Skip to content

[Bug]Add sequence_parallel in layernorm init to enable 3D parallelism with DeepSpeed for non CUDA device. #468

New issue

Have a question about this project? Sign up for a free GitHub account to open an issue and contact its maintainers and the community.

By clicking “Sign up for GitHub”, you agree to our terms of service and privacy statement. We’ll occasionally send you account related emails.

Already on GitHub? Sign in to your account

Open
wants to merge 1 commit into
base: main
Choose a base branch
from
Open
Show file tree
Hide file tree
Changes from all commits
Commits
File filter

Filter by extension

Filter by extension

Conversations
Failed to load comments.
Loading
Jump to
Jump to file
Failed to load files.
Loading
Diff view
Diff view
2 changes: 1 addition & 1 deletion megatron/model/__init__.py
Original file line number Diff line number Diff line change
Expand Up @@ -13,7 +13,7 @@
from .fused_rmsnorm import RMSNorm
else:
from .rmsnorm import RMSNorm
from torch.nn import LayerNorm
from .layernorm import LayerNorm

from .distributed import DistributedDataParallel
from .bert_model import BertModel
Expand Down
21 changes: 21 additions & 0 deletions megatron/model/layernorm.py
Original file line number Diff line number Diff line change
@@ -0,0 +1,21 @@
import numbers

import torch
from torch.nn.parameter import Parameter

class LayerNorm(torch.nn.Module):
def __init__(self, normalized_shape, eps: float = 1e-5, sequence_parallel=False):
super(LayerNorm, self).__init__()

if isinstance(normalized_shape, numbers.Integral):
normalized_shape = (normalized_shape,)
self.normalized_shape = torch.Size(normalized_shape)
self.eps = eps
self.weight = Parameter(torch.ones(normalized_shape))
self.bias = Parameter(torch.zeros(normalized_shape))
self.sequence_parallel = sequence_parallel
setattr(self.weight, 'sequence_parallel', self.sequence_parallel)

def forward(self, x):
output = torch.nn.functional.layer_norm(x, self.normalized_shape, self.weight, self.bias, self.eps)
return output