Skip to content

Commit 38ef10b

Browse files
albertzZettelkasten
andcommitted
CumConcatLayer
Co-authored-by: Frithjof <[email protected]>
1 parent 6f0cf76 commit 38ef10b

File tree

1 file changed

+165
-0
lines changed

1 file changed

+165
-0
lines changed

returnn/tf/layers/rec.py

+165
Original file line numberDiff line numberDiff line change
@@ -8495,3 +8495,168 @@ def get_out_data_from_opts(cls, name, sources, n_out, **kwargs):
84958495
kind=DimensionTag.Types.Spatial, description="%s_rel_pos_enc_time" % name, dimension=None)
84968496
data = data.copy_template_new_dim_tags((dummy_dim_tag, time_dim_tag, feature_dim_tag))
84978497
return data
8498+
8499+
8500+
class CumConcatLayer(_ConcatInputLayer):
8501+
"""
8502+
Concatenates all previous frames of a time-axis.
8503+
Like :class:`CumsumLayer` uses `sum`, this layer uses `concat`.
8504+
8505+
This layer expects to be inside a :class:`RecLayer`.
8506+
8507+
Inside a rec loop (not optimized out),
8508+
this will concatenate the current input
8509+
to the previous accumulated inputs.
8510+
For an input of shape `input_shape`,
8511+
it will output a tensor of shape `[new_dim] + input_shape`.
8512+
`new_dim` is a special dimension, usually of length `i`,
8513+
where `i` is the current loop frame,
8514+
i.e. the length increases in every loop frame.
8515+
`new_dim` is specified by a separate own dim tag.
8516+
For example, in the first frame,
8517+
this will be of shape `[1] + input_shape`,
8518+
in the second frame shape `[2] + input_shape`,
8519+
and so on,
8520+
and in the last frame shape `[T] + input_shape`.
8521+
8522+
Outside the rec loop (optimized out),
8523+
this layer expects an input with the time dim of the rec layer,
8524+
and returns the input as-is,
8525+
but replacing the time dim tag with the dim tag `new_dim`
8526+
converted as outside the loop.
8527+
8528+
Normally the optimization should not matter for the user,
8529+
i.e. for the user, the logical behavior is always as being inside the rec loop.
8530+
Outside the loop,
8531+
the output represents a tensor of shape `[T, new_dim] + input_shape`,
8532+
although we actually have another `new_dim` outside the loop,
8533+
and `T` is not actually there,
8534+
but we still have all the information,
8535+
because the last frame has all information.
8536+
"""
8537+
layer_class = "cum_concat"
8538+
recurrent = True # order matters
8539+
8540+
def __init__(self, new_dim, **kwargs):
8541+
"""
8542+
:param DimensionTag new_dim:
8543+
"""
8544+
super(CumConcatLayer, self).__init__(**kwargs)
8545+
rec_layer = self.network.get_rec_parent_layer(inside_loop=False)
8546+
assert rec_layer, "%r must be used inside a RecLayer" % self
8547+
out_axis = self.output.get_axis_from_description(new_dim)
8548+
new_dim_ = self.output.dim_tags[out_axis]
8549+
8550+
if self.network.is_inside_rec_layer(inside_loop=True):
8551+
current_data = self.input_data.copy_compatible_to(self.output, unbroadcast=False)
8552+
current_frame = current_data.placeholder # [B, 1, ..., D]
8553+
last_frames = self._rec_previous_layer.rec_vars_outputs["state"] # [B, t, ..., D]
8554+
concat_frames = tf.concat([last_frames, current_frame], axis=out_axis) # [B, t+1, ..., D]
8555+
self.rec_vars_outputs["state"] = concat_frames
8556+
self.output.placeholder = concat_frames
8557+
8558+
if not new_dim_.dyn_size_ext:
8559+
# Unbroadcasting to [B] is not needed because any layers operating on this
8560+
# should be able to handle extended dyn sizes.
8561+
# Clipping it to the max length for sequences in the loop which are already ended
8562+
# (i.e. considering the end flag)
8563+
# is also not needed because any calculations after the end are irrelevant.
8564+
dyn_size = self.network.get_rec_step_index() + 1 # scalar
8565+
new_dim_.dyn_size_ext = Data(
8566+
name="%s:cum-concat:size-inside" % self.name,
8567+
dim_tags=[], # scalar
8568+
placeholder=dyn_size)
8569+
8570+
else:
8571+
# If not inside a rec loop, this layer is a no-op on the tensor.
8572+
self.output.placeholder = self.input_data.placeholder
8573+
8574+
# However, we used new dim tags, which were already prepared.
8575+
# We now must fill in the extended dynamic size information.
8576+
if not new_dim_.dyn_size_ext:
8577+
# This must match the logic above for inside the loop.
8578+
dyn_size = tf.range(tf.math.reduce_max(rec_layer.time_dim_tag.dyn_size)) + 1 # [T]
8579+
new_dim_.dyn_size_ext = Data(
8580+
name="%s:cum-concat:size-outside" % self.name,
8581+
dim_tags=[rec_layer.time_dim_tag],
8582+
placeholder=dyn_size)
8583+
8584+
@classmethod
8585+
def get_out_data_from_opts(cls, name, network, sources, new_dim, **kwargs):
8586+
"""
8587+
:param str name:
8588+
:param returnn.tf.network.TFNetwork network:
8589+
:param list[LayerBase] sources:
8590+
:param DimensionTag new_dim:
8591+
:rtype: Data
8592+
"""
8593+
rec_layer = network.get_rec_parent_layer(inside_loop=False)
8594+
assert rec_layer, "CumConcatLayer %r must be used inside a RecLayer" % name
8595+
new_dim_base = new_dim.get_same_base()
8596+
if new_dim_base.per_spatial_frame is None:
8597+
new_dim_base.per_spatial_frame = rec_layer.time_dim_tag
8598+
else:
8599+
assert new_dim_base.per_spatial_frame == rec_layer.time_dim_tag
8600+
8601+
input_data = get_concat_sources_data_template(sources, name="%s_output" % name)
8602+
if network.is_inside_rec_layer(inside_loop=True):
8603+
# Currently SelectSearchSourcesLayer assumes that all rec_vars_outputs are batch-major.
8604+
# Therefore we here copy the input as batch-major, and then add the time axis at axis 1.
8605+
# In the future, when SelectSearchSourcesLayer has support for this, we can change this to operate on axis 0,
8606+
# which should be more efficient
8607+
out = input_data.copy_as_batch_major()
8608+
out = out.copy_add_dim_by_tag(new_dim_base, unbroadcast=True, axis=1)
8609+
return out
8610+
8611+
else: # outside loop
8612+
if not new_dim_base.per_spatial_frame_accumulated:
8613+
new_dim_accum = DimensionTag(
8614+
kind=new_dim_base.kind, description="%s:accumulated" % name)
8615+
new_dim_accum.same_as = new_dim_base
8616+
new_dim_base.per_spatial_frame_accumulated = new_dim_accum
8617+
else:
8618+
new_dim_accum = new_dim_base.per_spatial_frame_accumulated
8619+
# Assume that the input has the time dim from the rec layer.
8620+
axis = input_data.get_axis_from_description(rec_layer.time_dim_tag)
8621+
return input_data.copy_template_replace_dim_tag(axis=axis, new_dim_tag=new_dim_accum)
8622+
8623+
# noinspection PyMethodOverriding
8624+
@classmethod
8625+
def get_rec_initial_extra_outputs(cls, network, batch_dim, rec_layer, sources, output, new_dim, **kwargs):
8626+
"""
8627+
:param returnn.tf.network.TFNetwork network:
8628+
:param tf.Tensor batch_dim:
8629+
:param TFNetworkRecLayer.RecLayer|LayerBase rec_layer:
8630+
:param list[LayerBase] sources:
8631+
:param Data output:
8632+
:param DimensionTag new_dim:
8633+
:rtype: dict[str,tf.Tensor]
8634+
"""
8635+
if network.is_inside_rec_layer():
8636+
shape = []
8637+
for tag in output.dim_tags:
8638+
if tag.is_batch_dim():
8639+
shape.append(batch_dim)
8640+
elif tag == new_dim:
8641+
shape.append(0)
8642+
elif tag.dimension is not None:
8643+
shape.append(tag.dimension)
8644+
else:
8645+
assert tag.dyn_size is not None
8646+
shape.append(tf.math.reduce_max(tag.dyn_size))
8647+
return {"state": tf.zeros(shape, dtype=output.dtype)}
8648+
else:
8649+
return {}
8650+
8651+
@classmethod
8652+
def get_rec_initial_extra_outputs_shape_invariants(cls, network, sources, output, **kwargs):
8653+
"""
8654+
:param returnn.tf.network.TFNetwork network:
8655+
:param list[LayerBase] sources:
8656+
:param Data output:
8657+
:rtype: dict[str, tf.TensorShape]
8658+
"""
8659+
if network.is_inside_rec_layer():
8660+
return {"state": tf.TensorShape(output.batch_shape)}
8661+
else:
8662+
return {}

0 commit comments

Comments
 (0)