Skip to content

Commit

Permalink
add post_ln option
Browse files Browse the repository at this point in the history
  • Loading branch information
hpasapp committed Dec 15, 2020
1 parent 1569f31 commit e9d6ed2
Showing 1 changed file with 20 additions and 6 deletions.
26 changes: 20 additions & 6 deletions sru/modules.py
Original file line number Diff line number Diff line change
Expand Up @@ -41,7 +41,8 @@ def __init__(self,
rescale: bool = True,
v1: bool = False,
custom_m: Optional[nn.Module] = None,
amp_recurrence_fp16: bool = False):
amp_recurrence_fp16: bool = False,
post_ln: bool = False):
"""Initialize the SRUCell module.
Parameters
Expand Down Expand Up @@ -94,6 +95,8 @@ def __init__(self,
When using AMP autocast, selects which type to use
for recurrence custom kernel.
False: torch.float32, True: torch.float16
post_ln: bool
if True use post layer norm, else pre layer norm
"""
super(SRUCell, self).__init__()
self.input_size = input_size
Expand All @@ -113,6 +116,7 @@ def __init__(self,
self.activation_type = 1
self.activation = 'tanh'
self.amp_recurrence_fp16 = amp_recurrence_fp16
self.post_ln = post_ln

# projection dimension
self.projection_size = 0
Expand Down Expand Up @@ -146,7 +150,10 @@ def __init__(self,

self.layer_norm: Optional[nn.Module]= None
if layer_norm:
self.layer_norm = nn.LayerNorm(self.input_size)
if post_ln:
self.layer_norm = nn.LayerNorm(self.input_size)
else:
self.layer_norm = nn.LayerNorm(self.output_size)

self.reset_parameters()

Expand Down Expand Up @@ -235,7 +242,7 @@ def forward(self,

# apply layer norm before activation (i.e. before SRU computation)
residual = input
if self.layer_norm is not None:
if self.layer_norm is not None and not self.post_ln:
input = self.layer_norm(input)

# apply dropout for multiplication
Expand All @@ -260,6 +267,10 @@ def forward(self,

# apply elementwise recurrence to get hidden states h and c
h, c = self.apply_recurrence(U, V, residual, c0, scale_val, mask_c, mask_pad)

if self.layer_norm is not None and self.post_ln:
h = self.layer_norm(h)

return h, c

def apply_recurrence(self,
Expand Down Expand Up @@ -427,7 +438,8 @@ def __init__(self,
nn_rnn_compatible_return: bool = False,
custom_m: Optional[Union[nn.Module, List[nn.Module]]] = None,
proj_input_to_hidden_first: bool = False,
amp_recurrence_fp16: bool = False):
amp_recurrence_fp16: bool = False,
post_ln: bool = False):
"""Initialize the SRU module.
Parameters
Expand Down Expand Up @@ -486,7 +498,8 @@ def __init__(self,
When using AMP autocast, selects which type to use
for recurrence custom kernel.
False: torch.float32, True: torch.float16
post_ln: bool
if True use post layer norm, else use pre layer norm
"""

super(SRU, self).__init__()
Expand Down Expand Up @@ -538,7 +551,8 @@ def __init__(self,
rescale=rescale,
v1=v1,
custom_m=custom_m_i,
amp_recurrence_fp16=amp_recurrence_fp16
amp_recurrence_fp16=amp_recurrence_fp16,
post_ln=post_ln
)
rnn_lst.append(layer_i)
self.rnn_lst = rnn_lst
Expand Down

0 comments on commit e9d6ed2

Please sign in to comment.