def apply(self, inputs_q, inputs_kv, num_heads, dtype=jnp.float32, qkv_features=None, out_features=None, attention_axis=None, causal_mask=False, padding_mask=None, key_padding_mask=None, segmentation=None, key_segmentation=None, cache=None, broadcast_dropout=True, dropout_rng=None, dropout_rate=0., deterministic=False, precision=None, kernel_init=nn.linear.default_kernel_init, bias_init=nn.initializers.zeros, bias=True, block_size=50, max_num_blocks=25, sort_activation='softmax'): """Applies multi-head sinkhorn attention on the input data. Projects the inputs into multi-headed query, key, and value vectors, applies dot-product attention and project the results to an output vector. This can be used for encoder-decoder attention by specifying both `inputs_q` and `inputs_kv` orfor self-attention by only specifying `inputs_q` and setting `inputs_kv` to None. Args: inputs_q: input queries of shape `[bs, dim1, dim2, ..., dimN, features]`. inputs_kv: key/values of shape `[bs, dim1, dim2, ..., dimN, features]` or None for self-attention, inn which case key/values will be derived from inputs_q. num_heads: number of attention heads. Features (i.e. inputs_q.shape[-1]) should be divisible by the number of heads. dtype: the dtype of the computation (default: float32) qkv_features: dimension of the key, query, and value. out_features: dimension of the last projection attention_axis: axes over which the attention is applied ( 'None' means attention over all axes, but batch, heads, and features). causal_mask: boolean specifying whether to apply a causal mask on the attention weights. If True, the output at timestep `t` will not depend on inputs at timesteps strictly greater than `t`. padding_mask: boolean specifying query tokens that are pad token. key_padding_mask: boolean specifying key-value tokens that are pad token. segmentation: segment indices for packed inputs_q data. key_segmentation: segment indices for packed inputs_kv data. cache: an instance of `flax.nn.attention.Cache` used for efficient autoregressive decoding. broadcast_dropout: bool: use a broadcasted dropout along batch dims. dropout_rng: JAX PRNGKey: to be used for dropout dropout_rate: dropout rate deterministic: bool, deterministic or not (to apply dropout) precision: numerical precision of the computation see `jax.lax.Precision` for details. kernel_init: initializer for the kernel of the Dense layers. bias_init: initializer for the bias of the Dense layers. bias: bool: whether pointwise QKVO dense transforms use bias. block_size: int, block size. max_num_blocks: int, max num blocks. sort_activation: str {softmax, sinkhorn, gumbel_sinkhorn} Returns: output of shape `[bs, dim1, dim2, ..., dimN, features]`. """ assert causal_mask or not cache, ( 'Caching is only support for causal attention.') assert inputs_q.ndim == 3 if inputs_kv is None: inputs_kv = inputs_q if attention_axis is None: attention_axis = tuple(range(1, inputs_q.ndim - 1)) features = out_features or inputs_q.shape[-1] qkv_features = qkv_features or inputs_q.shape[-1] assert qkv_features % num_heads == 0, ( 'Memory dimension must be divisible by number of heads.') head_dim = qkv_features // num_heads dense = nn.DenseGeneral.partial( axis=-1, features=(num_heads, head_dim), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision) # project inputs_q to multi-headed q/k/v # dimensions are then [bs, dims..., n_heads, n_features_per_head] qlength = inputs_q.shape[-2] bs = inputs_q.shape[0] kvlength = inputs_kv.shape[-2] query, key, value = (dense(inputs_q, dtype=dtype, name='query'), dense(inputs_kv, dtype=dtype, name='key'), dense(inputs_kv, dtype=dtype, name='value')) if cache: assert isinstance(cache, Cache), 'cache must be an instance of Cache' if self.is_initializing(): cache.store(onp.array((key.ndim,) + key.shape[-2:], dtype=onp.int32)) else: cache_entry = cache.retrieve(None) expected_shape = list(cache_entry.key.shape[:-2]) for attn_dim in attention_axis: expected_shape[attn_dim] = 1 expected_shape = tuple(expected_shape) + inputs_q.shape[-1:] if expected_shape != inputs_q.shape: raise ValueError('Invalid shape provided, ' 'expected shape %s instead got %s.' % (expected_shape, inputs_q.shape)) if not isinstance(cache_entry, _CacheEntry): raise ValueError('Cache is not initialized.') cshape = cache_entry.key.shape indices = [0] * len(cshape) i = cache_entry.i attn_size = onp.prod(onp.take(cshape, attention_axis)) for attn_dim in attention_axis: attn_size //= cshape[attn_dim] indices[attn_dim] = i // attn_size i = i % attn_size key = lax.dynamic_update_slice(cache_entry.key, key, indices) value = lax.dynamic_update_slice(cache_entry.value, value, indices) one = jnp.array(1, jnp.uint32) cache_entry = cache_entry.replace(i=cache_entry.i + one, key=key, value=value) cache.store(cache_entry) key_padding_mask = jnp.broadcast_to( (jnp.arange(cshape[1]) < cache_entry.i), cshape[:2]) key_padding_mask = key_padding_mask.astype(jnp.float32)[..., None] # block reshape before attention num_query_blocks = qlength // block_size num_kv_blocks = kvlength // block_size block_query = jnp.reshape( query, (bs, block_size, num_query_blocks, num_heads, head_dim)) block_key = jnp.reshape( key, (bs, block_size, num_kv_blocks, num_heads, head_dim)) block_value = jnp.reshape( value, (bs, block_size, num_kv_blocks, num_heads, head_dim)) if causal_mask: # causal masking needs to not have blocks with mixed information. sum_key = jnp.cumsum(block_key, axis=1) sum_key = sum_key[:, 0, :, :, :] # take first item else: sum_key = jnp.sum(block_key, axis=1) # sort net on head_dim dimensions sort_out = nn.DenseGeneral(sum_key, axis=-1, features=(max_num_blocks), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision) # (bs x num_key_blocks x num_heads x num_key_blocks sort_out = sort_out[:, :, :, :num_query_blocks] # simple softmax sorting first. if sort_activation == 'sinkhorn': permutation = sinkhorn_operator( jnp.reshape(sort_out, (-1, num_kv_blocks, num_query_blocks)), causal=causal_mask) permutation = jnp.reshape(permutation, (-1, num_kv_blocks, num_heads, num_query_blocks)) else: if causal_mask: block_mask = _make_causal_mask(key, attention_axis) sort_out += block_mask permutation = jax.nn.softmax(sort_out, axis=-1) sorted_key = jnp.einsum('bskhd,bnhl->bsnhd', block_key, permutation) sorted_value = jnp.einsum('bskhd,bnhl->bsnhd', block_value, permutation) # create attention masks mask_components = [] sorted_mask_components = [] if causal_mask: # TODO(yitay): Test this causal masking. if cache and not self.is_initializing(): bias_pre_shape = (1,) * (key.ndim - 1) attn_shape = tuple(onp.take(key.shape, attention_axis)) attn_size = onp.prod(attn_shape) ii = jnp.arange(attn_size, dtype=jnp.uint32) mask = ii < cache_entry.i mask_components.append(mask.reshape(bias_pre_shape + attn_shape)) else: mask_components.append(_make_causal_mask(key, attention_axis)) if padding_mask is not None: # divide padding mask into block padding_mask = jnp.reshape(padding_mask, (bs * num_query_blocks, block_size, 1)) if key_padding_mask is None: key_padding_mask = padding_mask padding_mask = make_padding_mask( padding_mask_query=padding_mask, padding_mask_key=key_padding_mask, query_shape=(bs * num_query_blocks, block_size, num_heads, head_dim), key_shape=(bs * num_kv_blocks, block_size, num_heads, head_dim), attention_axis=attention_axis) padding_mask = jnp.reshape(padding_mask, (bs, num_query_blocks, block_size, block_size)) mask_components.append(padding_mask) sorted_padding_mask = jnp.einsum('bksj,bnhl->bnsj', padding_mask, permutation) sorted_mask_components.append(sorted_padding_mask) if segmentation is not None: if key_segmentation is None: key_segmentation = segmentation segmentation_mask = make_padding_mask( padding_mask_query=segmentation, padding_mask_key=key_segmentation, query_shape=(bs * num_query_blocks, block_size, num_heads, head_dim), key_shape=(bs * num_kv_blocks, block_size, num_heads, head_dim), attention_axis=attention_axis, segmentation_mask=True) segmentation_mask = jnp.reshape(segmentation_mask, (bs, num_query_blocks, block_size, block_size)) mask_components.append(segmentation_mask) sorted_segmentation_mask = jnp.einsum('bksj,bnhl->bnsj', segmentation_mask, permutation) sorted_mask_components.append(sorted_segmentation_mask) if mask_components: attention_mask = mask_components[0] for component in mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) # attention mask in the form of attention bias attention_bias = lax.select( attention_mask > 0, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: attention_bias = None if sorted_mask_components: attention_mask = sorted_mask_components[0] for component in sorted_mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) # attention mask in the form of attention bias sorted_attention_bias = lax.select( attention_mask > 0, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: sorted_attention_bias = None # apply attention x = local_dot_product_attention( block_query, block_key, block_value, dtype=dtype, axis=attention_axis, bias=attention_bias, precision=precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic) sorted_x = local_dot_product_attention( block_query, sorted_key, sorted_value, dtype=dtype, axis=attention_axis, bias=sorted_attention_bias, precision=precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic) x = x + sorted_x x = jnp.reshape(x, (bs, qlength, num_heads, head_dim)) # back to the original inputs dimensions out = nn.DenseGeneral( x, features=features, axis=(-2, -1), kernel_init=kernel_init, bias_init=bias_init, bias=bias, dtype=dtype, precision=precision, name='out') return out
def apply(self, inputs_q, inputs_kv, num_heads, dtype=jnp.float32, qkv_features=None, out_features=None, attention_axis=None, causal_mask=False, padding_mask=None, key_padding_mask=None, segmentation=None, key_segmentation=None, cache=None, broadcast_dropout=True, dropout_rng=None, dropout_rate=0., deterministic=False, precision=None, kernel_init=nn.linear.default_kernel_init, bias_init=nn.initializers.zeros, bias=True, num_partitions=2): """Applies multi-head dot product attention on the input data. Projects the inputs into multi-headed query, key, and value vectors, applies dot-product attention and project the results to an output vector. This can be used for encoder-decoder attention by specifying both `inputs_q` and `inputs_kv` orfor self-attention by only specifying `inputs_q` and setting `inputs_kv` to None. Args: inputs_q: input queries of shape `[bs, dim1, dim2, ..., dimN, features]`. inputs_kv: key/values of shape `[bs, dim1, dim2, ..., dimN, features]` or None for self-attention, inn which case key/values will be derived from inputs_q. num_heads: number of attention heads. Features (i.e. inputs_q.shape[-1]) should be divisible by the number of heads. dtype: the dtype of the computation (default: float32) qkv_features: dimension of the key, query, and value. out_features: dimension of the last projection attention_axis: axes over which the attention is applied ( 'None' means attention over all axes, but batch, heads, and features). causal_mask: boolean specifying whether to apply a causal mask on the attention weights. If True, the output at timestep `t` will not depend on inputs at timesteps strictly greater than `t`. padding_mask: boolean specifying query tokens that are pad token. key_padding_mask: boolean specifying key-value tokens that are pad token. segmentation: segment indices for packed inputs_q data. key_segmentation: segment indices for packed inputs_kv data. cache: an instance of `flax.nn.attention.Cache` used for efficient autoregressive decoding. broadcast_dropout: bool: use a broadcasted dropout along batch dims. dropout_rng: JAX PRNGKey: to be used for dropout dropout_rate: dropout rate deterministic: bool, deterministic or not (to apply dropout) precision: numerical precision of the computation see `jax.lax.Precision` for details. kernel_init: initializer for the kernel of the Dense layers. bias_init: initializer for the bias of the Dense layers. bias: bool: whether pointwise QKVO dense transforms use bias. num_partitions: number of ways to partition (i.e. how many devices to run across). Returns: output of shape `[bs, dim1, dim2, ..., dimN, features]`. """ assert causal_mask or not cache, ( 'Caching is only support for causal attention.') if inputs_kv is None: inputs_kv = inputs_q if attention_axis is None: attention_axis = tuple(range(1, inputs_q.ndim - 1)) features = out_features or inputs_q.shape[-1] qkv_features = qkv_features or inputs_q.shape[-1] assert qkv_features % num_heads == 0, ( 'Memory dimension must be divisible by number of heads.') head_dim = qkv_features // num_heads dense = nn.DenseGeneral.partial(axis=-1, features=(num_heads, head_dim), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision) # project inputs_q to multi-headed q/k/v # dimensions are then [bs, dims..., n_heads, n_features_per_head] query, key, value = (dense(inputs_q, dtype=dtype, name='query'), dense(inputs_kv, dtype=dtype, name='key'), dense(inputs_kv, dtype=dtype, name='value')) if num_partitions > 1: partitions = P(1, 1, num_partitions, 1) query = with_sharding_constraint(query, partitions) key = with_sharding_constraint(key, partitions) value = with_sharding_constraint(value, partitions) if cache: assert isinstance(cache, Cache), 'cache must be an instance of Cache' if self.is_initializing(): cache.store(lambda: (key.ndim, key.shape[-2:])) else: cache_entry = cache.retrieve(None) expected_shape = list(cache_entry.key.shape[:-2]) for attn_dim in attention_axis: expected_shape[attn_dim] = 1 expected_shape = tuple(expected_shape) + inputs_q.shape[-1:] if expected_shape != inputs_q.shape: raise ValueError('Invalid shape provided, ' 'expected shape %s instead got %s.' % (expected_shape, inputs_q.shape)) if not isinstance(cache_entry, _CacheEntry): raise ValueError('Cache is not initialized.') cshape = cache_entry.key.shape i = cache_entry.i one_hot_indices = jax.nn.one_hot(i, cshape[3], dtype=key.dtype).reshape( (1, 1, 1, cshape[3])) key = key.transpose((0, 2, 3, 1)) key = cache_entry.key + key * one_hot_indices value = value.transpose((0, 2, 3, 1)) value = cache_entry.value + value * one_hot_indices one = jnp.array(1, jnp.uint32) cache_entry = cache_entry.replace(i=cache_entry.i + one, key=key, value=value) cache.store(cache_entry) key = key.transpose((0, 3, 1, 2)) value = value.transpose((0, 3, 1, 2)) cshape = (cshape[0], cshape[3], cshape[1], cshape[2]) # TODO(levskaya): verify this is still needed in translation decoding. key_padding_mask = jnp.broadcast_to( (jnp.arange(cshape[1]) < cache_entry.i), cshape[:2]) key_padding_mask = key_padding_mask.astype(jnp.float32)[..., None] # create attention masks mask_components = [] if causal_mask: if cache and not self.is_initializing(): bias_pre_shape = (1, ) * (key.ndim - 1) attn_shape = tuple(np.take(key.shape, attention_axis)) attn_size = np.prod(attn_shape) ii = jnp.arange(attn_size, dtype=jnp.uint32) mask = ii < cache_entry.i mask_components.append( mask.reshape(bias_pre_shape + attn_shape)) else: mask_components.append(_make_causal_mask(key, attention_axis)) if padding_mask is not None: if key_padding_mask is None: key_padding_mask = padding_mask padding_mask = make_padding_mask(padding_mask_query=padding_mask, padding_mask_key=key_padding_mask, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis) mask_components.append(padding_mask) if segmentation is not None: if key_segmentation is None: key_segmentation = segmentation segmentation_mask = make_padding_mask( padding_mask_query=segmentation, padding_mask_key=key_segmentation, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis, segmentation_mask=True) mask_components.append(segmentation_mask) if mask_components: attention_mask = mask_components[0] for component in mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) # attention mask in the form of attention bias attention_bias = lax.select( attention_mask > 0, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: attention_bias = None # apply attention x = dot_product_attention(query, key, value, dtype=dtype, axis=attention_axis, bias=attention_bias, precision=precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic) # back to the original inputs dimensions out = nn.DenseGeneral(x, features=features, axis=(-2, -1), kernel_init=kernel_init, bias_init=bias_init, bias=bias, dtype=dtype, precision=precision, name='out') if num_partitions > 1: x = with_sharding_constraint(x, None) return out
def __call__(self, inputs_q, inputs_kv, *, padding_mask, key_padding_mask, segmentation=None, key_segmentation=None): """Applies multi-head dot product attention on the input data. If weight_prec is not None, scales and quantizes weights to signed int with weight_prec bits. Projects the inputs into multi-headed query, key, and value vectors, applies dot-product attention and project the results to an output vector. This can be used for encoder-decoder attention by specifying both `inputs_q` and `inputs_kv` or for self-attention by only specifying `inputs_q` and setting `inputs_kv` to None. Args: inputs_q: input queries of shape `[bs, dim1, dim2, ..., dimN, features]`. inputs_kv: key/values of shape `[bs, dim1, dim2, ..., dimN, features]` or None for self-attention, inn which case key/values will be derived from inputs_q. padding_mask: boolean tensor specifying query tokens that are pad token. key_padding_mask: boolean tensor specifying key-value tokens that are pad token. segmentation: segment indices for packed inputs_q data. key_segmentation: segment indices for packed inputs_kv data. Returns: output of shape `[bs, dim1, dim2, ..., dimN, features]`. """ batch_size, query_sequence_length, channel_size = inputs_q.shape hparams = self.hparams if inputs_kv is None: inputs_kv = inputs_q key_sequence_length = inputs_q.shape[1] else: key_sequence_length = inputs_kv.shape[1] shape_utils.assert_shapes_equal( inputs_kv.shape, (batch_size, key_sequence_length, channel_size)) jax_precision = jax.lax.Precision.DEFAULT if padding_mask is not None: shape_utils.assert_shapes_equal( padding_mask.shape, (batch_size, query_sequence_length, 1)) if key_padding_mask is None: key_padding_mask = padding_mask else: shape_utils.assert_shapes_equal( key_padding_mask.shape, (batch_size, key_sequence_length, 1)) attention_axis = self.attention_axis if attention_axis is None: attention_axis = tuple(range(1, inputs_q.ndim - 1)) qkv_features = self.qkv_features qkv_features = qkv_features or inputs_q.shape[-1] num_heads = self.num_heads assert qkv_features % num_heads == 0, ( 'Memory dimension must be divisible by number of heads.') head_dim = qkv_features // num_heads paxis_name = self.paxis_name train = self.train kernel_init = self.kernel_init bias_init = self.bias_init use_bias = self.use_bias dtype = self.dtype def multi_batch_dense_aqt(inputs, *, name, padding_mask): batch_size, sequence_length, channel_size = inputs.shape inputs = inputs.reshape(batch_size * sequence_length, channel_size) if padding_mask is not None: padding_mask = padding_mask.reshape( batch_size * sequence_length, 1) out = flax_layers.DenseAqt(name=name, features=num_heads * head_dim, paxis_name=paxis_name, train=train, quant_context=self.quant_context, hparams=hparams.dense_kqv, kernel_init=kernel_init, bias_init=bias_init, use_bias=use_bias, dtype=dtype)(inputs, padding_mask=padding_mask) return out.reshape(batch_size, sequence_length, num_heads, head_dim) # project inputs_q to multi-headed q/k/v # dimensions are then [bs, sequence_length, n_heads, n_features_per_head] query = multi_batch_dense_aqt(inputs_q, name='query', padding_mask=padding_mask) key = multi_batch_dense_aqt(inputs_kv, name='key', padding_mask=key_padding_mask) value = multi_batch_dense_aqt(inputs_kv, name='value', padding_mask=key_padding_mask) is_cache_initialized = False if self.decode: is_cache_initialized = self.has_variable('cache', 'cached_key') cached_key = self.variable('cache', 'cached_key', jnp.zeros, key.shape, key.dtype) cached_value = self.variable('cache', 'cached_value', jnp.zeros, value.shape, value.dtype) cache_index = self.variable('cache', 'cache_index', lambda: jnp.array(0, dtype=jnp.int32)) if is_cache_initialized: expected_shape = list(cached_key.value.shape[:-2]) for attn_dim in attention_axis: expected_shape[attn_dim] = 1 expected_shape = tuple(expected_shape) + inputs_q.shape[-1:] if expected_shape != inputs_q.shape: raise ValueError('Invalid shape provided, ' 'expected shape %s instead got %s.' % (expected_shape, inputs_q.shape)) cshape = cached_key.value.shape indices = [0] * len(cshape) i = cache_index.value attn_size = onp.prod(onp.take(cshape, attention_axis)) *batch_dims, max_length, num_heads, depth_per_head = ( # pylint: disable=unused-variable cached_key.value.shape) indices = (0, ) * len(batch_dims) + (i, 0, 0) key = lax.dynamic_update_slice(cached_key.value, key, indices) value = lax.dynamic_update_slice(cached_value.value, value, indices) one = jnp.array(1, jnp.int32) cache_index.value = cache_index.value + one cached_key.value = key cached_value.value = value # TODO(levskaya): verify this is still needed in translation decoding. key_padding_mask = jnp.broadcast_to( (jnp.arange(max_length) < cache_index.value), cshape[:2]) key_padding_mask = key_padding_mask.astype( jnp.float32)[Ellipsis, None] # create attention masks mask_components = [] if self.causal_mask: if self.decode and is_cache_initialized: bias_pre_shape = (1, ) * (key.ndim - 1) attn_shape = tuple(onp.take(key.shape, attention_axis)) attn_size = onp.prod(attn_shape) ii = jnp.arange(attn_size, dtype=jnp.int32) mask = ii < cache_index.value mask_components.append( mask.reshape(bias_pre_shape + attn_shape)) else: mask_components.append(_make_causal_mask(key, attention_axis)) if padding_mask is not None: if key_padding_mask is None: key_padding_mask = padding_mask attn_padding_mask = make_padding_mask( padding_mask_query=padding_mask, padding_mask_key=key_padding_mask, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis) mask_components.append(attn_padding_mask) if segmentation is not None: if key_segmentation is None: key_segmentation = segmentation segmentation_mask = make_padding_mask( padding_mask_query=segmentation, padding_mask_key=key_segmentation, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis, segmentation_mask=True) mask_components.append(segmentation_mask) attention_mask = None if mask_components: attention_mask = mask_components[0] for component in mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) attention_mask = attention_mask.astype(jnp.bool_) # attention mask in the form of attention bias attention_bias = jnp.where( attention_mask, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: attention_bias = None # Add an extra dimension to the mask corresponding to the head # dimension. eg, if inputs_q has shape [batch_size, sequence_length, # n_features], then padding_mask will have a shape # [batch_size, sequence_length, 1] and query will have shape # [batch_size, sequence_length, n_heads, n_features_per_head]. # We create query_padding_mask with shape [batch_size, sequence_length, # 1, 1] to be broadcast-compatible with 'query'. if padding_mask is not None: padding_mask = padding_mask[Ellipsis, None] shape_utils.assert_shapes_equal( padding_mask.shape, (batch_size, query_sequence_length, 1, 1)) if key_padding_mask is not None: key_padding_mask = key_padding_mask[Ellipsis, None] # During prediction, the key padding mask is only going to be # broadcast-compatible with the key. shape_utils.assert_shapes_compatible( key_padding_mask.shape, (batch_size, key_sequence_length, 1, 1)) # apply attention attention_fn = self.attention_fn dropout_rate = self.dropout_rate broadcast_dropout = self.broadcast_dropout deterministic = self.deterministic if not deterministic and self.dropout_rate > 0.0: dropout_rng = self.make_rng('dropout') else: dropout_rng = None x = attention_fn( # pylint: disable=redundant-keyword-arg query=query, key=key, value=value, hparams=hparams.attn_acts, paxis_name=paxis_name, train=train, quant_context=self.quant_context, dtype=dtype, axis=attention_axis, bias=attention_bias, precision=jax_precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic, query_padding_mask=padding_mask, key_padding_mask=key_padding_mask, attn_mask=attention_mask) shape_utils.assert_shapes_equal( x.shape, (batch_size, query_sequence_length, num_heads, head_dim)) x = x.reshape(batch_size * query_sequence_length, num_heads * head_dim) if padding_mask is not None: padding_mask = padding_mask.reshape( batch_size * query_sequence_length, 1) # back to the original inputs dimensions out = flax_layers.DenseAqt(features=channel_size, hparams=hparams.dense_out, quant_context=self.quant_context, paxis_name=paxis_name, train=train, kernel_init=kernel_init, bias_init=bias_init, use_bias=use_bias, dtype=dtype, name='dense_out')(x, padding_mask=padding_mask) shape_utils.assert_shapes_equal( out.shape, (batch_size * query_sequence_length, channel_size)) out = out.reshape(batch_size, query_sequence_length, channel_size) return out
def apply(self, inputs_q, inputs_kv, num_heads, dtype=jnp.float32, qkv_features=None, out_features=None, attention_axis=None, causal_mask=False, padding_mask=None, key_padding_mask=None, segmentation=None, key_segmentation=None, cache=None, broadcast_dropout=True, dropout_rng=None, dropout_rate=0., deterministic=False, precision=None, kernel_init=nn.linear.default_kernel_init, bias_init=nn.initializers.zeros, bias=True, max_length=512, ignore_dot_product=True, synthesizer_mode='factorized_random', k=32): """Applies multi-head synthesizer attention on the input data. Projects the inputs into multi-headed query, key, and value vectors, applies dot-product attention and project the results to an output vector. This can be used for encoder-decoder attention by specifying both `inputs_q` and `inputs_kv` orfor self-attention by only specifying `inputs_q` and setting `inputs_kv` to None. Args: inputs_q: input queries of shape `[bs, dim1, dim2, ..., dimN, features]`. inputs_kv: key/values of shape `[bs, dim1, dim2, ..., dimN, features]` or None for self-attention, inn which case key/values will be derived from inputs_q. num_heads: number of attention heads. Features (i.e. inputs_q.shape[-1]) should be divisible by the number of heads. dtype: the dtype of the computation (default: float32) qkv_features: dimension of the key, query, and value. out_features: dimension of the last projection attention_axis: axes over which the attention is applied ( 'None' means attention over all axes, but batch, heads, and features). causal_mask: boolean specifying whether to apply a causal mask on the attention weights. If True, the output at timestep `t` will not depend on inputs at timesteps strictly greater than `t`. padding_mask: boolean specifying query tokens that are pad token. key_padding_mask: boolean specifying key-value tokens that are pad token. segmentation: segment indices for packed inputs_q data. key_segmentation: segment indices for packed inputs_kv data. cache: an instance of `flax.nn.attention.Cache` used for efficient autoregressive decoding. broadcast_dropout: bool: use a broadcasted dropout along batch dims. dropout_rng: JAX PRNGKey: to be used for dropout dropout_rate: dropout rate deterministic: bool, deterministic or not (to apply dropout) precision: numerical precision of the computation see `jax.lax.Precision` for details. kernel_init: initializer for the kernel of the Dense layers. bias_init: initializer for the bias of the Dense layers. bias: bool: whether pointwise QKVO dense transforms use bias. max_length: int, the maximum supported sequence length. ignore_dot_product: bool, to ignore the dot product attention or not. synthesizer_mode: str support 'dense' and 'random' or 'dense+random' k: int, low rank factorized attention. Returns: output of shape `[bs, dim1, dim2, ..., dimN, features]`. """ assert causal_mask or not cache, ( 'Caching is only support for causal attention.') assert inputs_q.ndim == 3 if inputs_kv is None: inputs_kv = inputs_q if attention_axis is None: attention_axis = tuple(range(1, inputs_q.ndim - 1)) features = out_features or inputs_q.shape[-1] qkv_features = qkv_features or inputs_q.shape[-1] assert qkv_features % num_heads == 0, ( 'Memory dimension must be divisible by number of heads.') head_dim = qkv_features // num_heads dense = nn.DenseGeneral.partial(axis=-1, features=(num_heads, head_dim), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision) # project inputs_q to multi-headed q/k/v # dimensions are then [bs, dims..., n_heads, n_features_per_head] qlength = inputs_q.shape[-2] kvlength = inputs_kv.shape[-2] if ignore_dot_product: value = dense(inputs_kv, dtype=dtype, name='value') key = value query = inputs_q else: query, key, value = (dense(inputs_q, dtype=dtype, name='query'), dense(inputs_kv, dtype=dtype, name='key'), dense(inputs_kv, dtype=dtype, name='value')) syn_weights_list = [] logging.info(synthesizer_mode) if 'random' in synthesizer_mode: if 'factorized_random' in synthesizer_mode: logging.info('Using factorized random') rand_syn_weights1 = self.param('random1', (num_heads, max_length, k), kernel_init) rand_syn_weights2 = self.param('random2', (num_heads, k, max_length), kernel_init) rand_syn_weights1 = rand_syn_weights1[:, :qlength, :] rand_syn_weights2 = rand_syn_weights2[:, :, :kvlength] rand_syn_weights = jnp.einsum('hlk,hkn->hln', rand_syn_weights1, rand_syn_weights2) rand_syn_weights = jax.lax.broadcast(rand_syn_weights, (inputs_q.shape[0], )) syn_weights_list.append(rand_syn_weights) else: rand_syn_weights = self.param( 'random', (num_heads, max_length, max_length), kernel_init) rand_syn_weights = rand_syn_weights[:, :qlength, :kvlength] rand_syn_weights = jax.lax.broadcast(rand_syn_weights, (inputs_q.shape[0], )) syn_weights_list.append(rand_syn_weights) if 'dense' in synthesizer_mode: dense_syn = nn.DenseGeneral.partial(axis=-1, features=(num_heads, head_dim), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision, name='dense_syn', dtype=dtype) # TODO(yitay): Change this to nn.Dense and make sure it works dense_syn_length = nn.linear.DenseGeneral.partial( axis=-1, features=(max_length), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision, name='dense_syn2', dtype=dtype) proj = dense_syn(inputs_q, dtype=dtype, name='dense_syn') proj = jax.nn.relu(proj) proj = dense_syn_length(proj, dtype=dtype, name='dense_syn_len') # TODO(yitay) check if this reshape is needed dense_syn_weights = proj.reshape( (inputs_q.shape[0], num_heads, qlength, max_length)) dense_syn_weights = dense_syn_weights[:, :, :, :qlength] syn_weights_list.append(dense_syn_weights) if cache: assert isinstance(cache, Cache), 'cache must be an instance of Cache' if self.is_initializing(): cache.store( onp.array((key.ndim, ) + key.shape[-2:], dtype=onp.int32)) else: cache_entry = cache.retrieve(None) expected_shape = list(cache_entry.key.shape[:-2]) for attn_dim in attention_axis: expected_shape[attn_dim] = 1 expected_shape = tuple(expected_shape) + inputs_q.shape[-1:] if expected_shape != inputs_q.shape: raise ValueError('Invalid shape provided, ' 'expected shape %s instead got %s.' % (expected_shape, inputs_q.shape)) if not isinstance(cache_entry, _CacheEntry): raise ValueError('Cache is not initialized.') cshape = cache_entry.key.shape indices = [0] * len(cshape) i = cache_entry.i attn_size = onp.prod(onp.take(cshape, attention_axis)) for attn_dim in attention_axis: attn_size //= cshape[attn_dim] indices[attn_dim] = i // attn_size i = i % attn_size key = lax.dynamic_update_slice(cache_entry.key, key, indices) value = lax.dynamic_update_slice(cache_entry.value, value, indices) one = jnp.array(1, jnp.uint32) cache_entry = cache_entry.replace(i=cache_entry.i + one, key=key, value=value) cache.store(cache_entry) key_padding_mask = jnp.broadcast_to( (jnp.arange(cshape[1]) < cache_entry.i), cshape[:2]) key_padding_mask = key_padding_mask.astype(jnp.float32)[..., None] # create attention masks mask_components = [] if causal_mask: if cache and not self.is_initializing(): bias_pre_shape = (1, ) * (key.ndim - 1) attn_shape = tuple(onp.take(key.shape, attention_axis)) attn_size = onp.prod(attn_shape) ii = jnp.arange(attn_size, dtype=jnp.uint32) mask = ii < cache_entry.i mask_components.append( mask.reshape(bias_pre_shape + attn_shape)) else: mask_components.append(_make_causal_mask(key, attention_axis)) if not ignore_dot_product: if padding_mask is not None: if key_padding_mask is None: key_padding_mask = padding_mask padding_mask = make_padding_mask( padding_mask_query=padding_mask, padding_mask_key=key_padding_mask, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis) mask_components.append(padding_mask) if segmentation is not None: if key_segmentation is None: key_segmentation = segmentation segmentation_mask = make_padding_mask( padding_mask_query=segmentation, padding_mask_key=key_segmentation, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis, segmentation_mask=True) mask_components.append(segmentation_mask) if mask_components: attention_mask = mask_components[0] for component in mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) # attention mask in the form of attention bias attention_bias = lax.select( attention_mask > 0, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: attention_bias = None # apply attention x = synthetic_attention(query, key, value, syn_weights_list, dtype=dtype, axis=attention_axis, bias=attention_bias, precision=precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic, ignore_dot_product=ignore_dot_product) # back to the original inputs dimensions out = nn.DenseGeneral(x, features=features, axis=(-2, -1), kernel_init=kernel_init, bias_init=bias_init, bias=bias, dtype=dtype, precision=precision, name='out') return out
def apply(self, inputs_q, inputs_kv, num_heads, block_size=64, num_rand_blocks=3, dtype=jnp.float32, qkv_features=None, out_features=None, attention_axis=None, causal_mask=False, padding_mask=None, key_padding_mask=None, segmentation=None, key_segmentation=None, cache=None, broadcast_dropout=True, dropout_rng=None, dropout_rate=0., deterministic=False, precision=None, kernel_init=nn.linear.default_kernel_init, bias_init=nn.initializers.zeros, bias=True, connectivity_seed=None): """Applies multi-head dot product attention on the input data. Projects the inputs into multi-headed query, key, and value vectors, applies dot-product attention and project the results to an output vector. This can be used for encoder-decoder attention by specifying both `inputs_q` and `inputs_kv` orfor self-attention by only specifying `inputs_q` and setting `inputs_kv` to None. Args: inputs_q: input queries of shape `[bs, length, features]`. inputs_kv: key/values of shape `[bs, length, features]` or None for self-attention, inn which case key/values will be derived from inputs_q. num_heads: number of attention heads. Features (i.e. inputs_q.shape[-1]) should be divisible by the number of heads. block_size: Size for local attention around diagonal of attention. num_rand_blocks: int. Number of random chunks per row. dtype: the dtype of the computation (default: float32) qkv_features: dimension of the key, query, and value. out_features: dimension of the last projection attention_axis: axes over which the attention is applied ( 'None' means attention over all axes, but batch, heads, and features). causal_mask: boolean specifying whether to apply a causal mask on the attention weights. If True, the output at timestep `t` will not depend on inputs at timesteps strictly greater than `t`. padding_mask: boolean specifying query tokens that are pad token. key_padding_mask: boolean specifying key-value tokens that are pad token. segmentation: segment indices for packed inputs_q data. key_segmentation: segment indices for packed inputs_kv data. cache: an instance of `flax.nn.attention.Cache` used for efficient autoregressive decoding. broadcast_dropout: bool: use a broadcasted dropout along batch dims. dropout_rng: JAX PRNGKey: to be used for dropout dropout_rate: dropout rate deterministic: bool, deterministic or not (to apply dropout) precision: numerical precision of the computation see `jax.lax.Precision` for details. kernel_init: initializer for the kernel of the Dense layers. bias_init: initializer for the bias of the Dense layers. bias: bool: whether pointwise QKVO dense transforms use bias. connectivity_seed: Seed for random block sparse attention. Returns: output of shape `[bs, length, features]`. """ orig_seqlen = inputs_q.shape[-2] logging.info(inputs_q) extra_len = block_size - (orig_seqlen % block_size) pad_width = jnp.array([[0, 0], [0, extra_len], [0, 0]]) mask_pad = jnp.array([[0, 0], [0, extra_len], [0, 0]]) padding_mask = jnp.pad(padding_mask, mask_pad, constant_values=-1e9) inputs_q = jnp.pad(inputs_q, pad_width) if inputs_kv is not None: inputs_kv = jnp.pad(inputs_kv, pad_width) assert causal_mask or not cache, ( 'Caching is only support for causal attention.') if inputs_kv is None: inputs_kv = inputs_q if attention_axis is None: attention_axis = tuple(range(1, inputs_q.ndim - 1)) features = out_features or inputs_q.shape[-1] qkv_features = qkv_features or inputs_q.shape[-1] assert qkv_features % num_heads == 0, ( 'Memory dimension must be divisible by number of heads.') head_dim = qkv_features // num_heads dense = nn.DenseGeneral.partial(axis=-1, features=(num_heads, head_dim), kernel_init=kernel_init, bias_init=bias_init, bias=bias, precision=precision) # project inputs_q to multi-headed q/k/v # dimensions are then [bs, dims..., n_heads, n_features_per_head] query, key, value = (dense(inputs_q, dtype=dtype, name='query'), dense(inputs_kv, dtype=dtype, name='key'), dense(inputs_kv, dtype=dtype, name='value')) if cache: assert isinstance( cache, attention.Cache), 'cache must be an instance of Cache' if self.is_initializing(): cache.store( onp.array((key.ndim, ) + key.shape[-2:], dtype=onp.int32)) else: cache_entry = cache.retrieve(None) expected_shape = list(cache_entry.key.shape[:-2]) for attn_dim in attention_axis: expected_shape[attn_dim] = 1 expected_shape = tuple(expected_shape) + inputs_q.shape[-1:] if expected_shape != inputs_q.shape: raise ValueError('Invalid shape provided, ' 'expected shape %s instead got %s.' % (expected_shape, inputs_q.shape)) if not isinstance(cache_entry, attention._CacheEntry): # pylint: disable=protected-access raise ValueError('Cache is not initialized.') cshape = cache_entry.key.shape indices = [0] * len(cshape) i = cache_entry.i attn_size = onp.prod(onp.take(cshape, attention_axis)) for attn_dim in attention_axis: attn_size //= cshape[attn_dim] indices[attn_dim] = i // attn_size i = i % attn_size key = lax.dynamic_update_slice(cache_entry.key, key, indices) value = lax.dynamic_update_slice(cache_entry.value, value, indices) one = jnp.array(1, jnp.uint32) cache_entry = cache_entry.replace(i=cache_entry.i + one, key=key, value=value) cache.store(cache_entry) # TODO(levskaya): verify this is still needed in translation decoding. key_padding_mask = jnp.broadcast_to( (jnp.arange(cshape[1]) < cache_entry.i), cshape[:2]) key_padding_mask = key_padding_mask.astype(jnp.float32)[..., None] if causal_mask: # Falls back to full attention with a causal mask. # create attention masks mask_components = [] if causal_mask: if cache and not self.is_initializing(): bias_pre_shape = (1, ) * (key.ndim - 1) attn_shape = tuple(onp.take(key.shape, attention_axis)) attn_size = onp.prod(attn_shape) ii = jnp.arange(attn_size, dtype=jnp.uint32) mask = ii < cache_entry.i mask_components.append( mask.reshape(bias_pre_shape + attn_shape)) else: mask_components.append( attention._make_causal_mask(key, attention_axis)) # pylint: disable=protected-access if padding_mask is not None: if key_padding_mask is None: key_padding_mask = padding_mask padding_mask = attention.make_padding_mask( padding_mask_query=padding_mask, padding_mask_key=key_padding_mask, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis) mask_components.append(padding_mask) if segmentation is not None: if key_segmentation is None: key_segmentation = segmentation segmentation_mask = attention.make_padding_mask( padding_mask_query=segmentation, padding_mask_key=key_segmentation, query_shape=query.shape, key_shape=key.shape, attention_axis=attention_axis, segmentation_mask=True) mask_components.append(segmentation_mask) if mask_components: attention_mask = mask_components[0] for component in mask_components[1:]: attention_mask = jnp.logical_and(attention_mask, component) # attention mask in the form of attention bias attention_bias = lax.select( attention_mask > 0, jnp.full(attention_mask.shape, 0.).astype(dtype), jnp.full(attention_mask.shape, -1e10).astype(dtype)) else: attention_bias = None x = nn.dot_product_attention(query, key, value, dtype=dtype, axis=attention_axis, bias=attention_bias, precision=precision, dropout_rng=dropout_rng, dropout_rate=dropout_rate, broadcast_dropout=broadcast_dropout, deterministic=deterministic) else: if connectivity_seed is None: path = self._get_construction_frame().path connectivity_seed = hash(path) % 2**32 # apply attention input_mask = None if padding_mask is not None: input_mask = padding_mask.astype(key.dtype) x = sparse_dot_product_attention( query, key, value, connectivity_seed=connectivity_seed, input_mask=input_mask, block_size=block_size, num_rand_blocks=num_rand_blocks) # back to the original inputs dimensions out = nn.DenseGeneral(x, features=features, axis=(-2, -1), kernel_init=kernel_init, bias_init=bias_init, bias=bias, dtype=dtype, precision=precision, name='out') out = out[:, :orig_seqlen, :] return out