mirror of
https://github.com/hpcaitech/ColossalAI.git
synced 2025-09-09 21:09:18 +00:00
[bf16] add bf16 support (#3882)
* [bf16] add bf16 support for fused adam (#3844) * [bf16] fused adam kernel support bf16 * [test] update fused adam kernel test * [test] update fused adam test * [bf16] cpu adam and hybrid adam optimizers support bf16 (#3860) * [bf16] implement mixed precision mixin and add bf16 support for low level zero (#3869) * [bf16] add mixed precision mixin * [bf16] low level zero optim support bf16 * [text] update low level zero test * [text] fix low level zero grad acc test * [bf16] add bf16 support for gemini (#3872) * [bf16] gemini support bf16 * [test] update gemini bf16 test * [doc] update gemini docstring * [bf16] add bf16 support for plugins (#3877) * [bf16] add bf16 support for legacy zero (#3879) * [zero] init context support bf16 * [zero] legacy zero support bf16 * [test] add zero bf16 test * [doc] add bf16 related docstring for legacy zero
This commit is contained in:
131
tests/test_optimizer/test_adam_kernel.py
Normal file
131
tests/test_optimizer/test_adam_kernel.py
Normal file
@@ -0,0 +1,131 @@
|
||||
# This test checks adam kernels
|
||||
# Baseline is pure fp32 torch adam optimizer
|
||||
import math
|
||||
from abc import abstractmethod
|
||||
from typing import Type
|
||||
|
||||
import pytest
|
||||
import torch
|
||||
from torch import Tensor
|
||||
|
||||
from colossalai.utils import get_current_device, multi_tensor_applier
|
||||
|
||||
_FUSED_ALLOWED_P_G_TYPES = [(torch.float, torch.half), (torch.float, torch.float), (torch.half, torch.float),
|
||||
(torch.half, torch.half), (torch.bfloat16, torch.float), (torch.float, torch.bfloat16),
|
||||
(torch.bfloat16, torch.bfloat16)]
|
||||
|
||||
_CPU_ALLOWED_P_G_TYPES = [(torch.float, torch.half), (torch.float, torch.float), (torch.half, torch.float),
|
||||
(torch.half, torch.half)]
|
||||
|
||||
|
||||
class AdamKernel:
|
||||
|
||||
def __init__(self, lr: float, beta1: float, beta2: float, eps: float, weight_decay: float, use_adamw: bool) -> None:
|
||||
self.lr = lr
|
||||
self.beta1 = beta1
|
||||
self.beta2 = beta2
|
||||
self.eps = eps
|
||||
self.weight_decay = weight_decay
|
||||
self.use_adamw = use_adamw
|
||||
|
||||
@abstractmethod
|
||||
def update(self, step: int, param: Tensor, grad: Tensor, exp_avg: Tensor, exp_avg_sq: Tensor):
|
||||
pass
|
||||
|
||||
|
||||
class TorchAdamKernel(AdamKernel):
|
||||
|
||||
def update(self, step: int, param: Tensor, grad: Tensor, exp_avg: Tensor, exp_avg_sq: Tensor):
|
||||
bias_correction1 = 1 - self.beta1**step
|
||||
bias_correction2 = 1 - self.beta2**step
|
||||
|
||||
if self.weight_decay != 0:
|
||||
if self.use_adamw:
|
||||
# Perform stepweight decay
|
||||
param.mul_(1 - self.lr * self.weight_decay)
|
||||
else:
|
||||
grad = grad.add(param, alpha=self.weight_decay)
|
||||
|
||||
# Decay the first and second moment running average coefficient
|
||||
exp_avg.mul_(self.beta1).add_(grad, alpha=1 - self.beta1)
|
||||
exp_avg_sq.mul_(self.beta2).addcmul_(grad, grad, value=1 - self.beta2)
|
||||
denom = (exp_avg_sq.sqrt() / math.sqrt(bias_correction2)).add_(self.eps)
|
||||
|
||||
step_size = self.lr / bias_correction1
|
||||
|
||||
param.addcdiv_(exp_avg, denom, value=-step_size)
|
||||
|
||||
|
||||
class FusedAdamKernel(AdamKernel):
|
||||
|
||||
def __init__(self, lr: float, beta1: float, beta2: float, eps: float, weight_decay: float, use_adamw: bool) -> None:
|
||||
super().__init__(lr, beta1, beta2, eps, weight_decay, use_adamw)
|
||||
from colossalai.kernel.op_builder import FusedOptimBuilder
|
||||
fused_optim = FusedOptimBuilder().load()
|
||||
self.fused_adam = fused_optim.multi_tensor_adam
|
||||
self.dummy_overflow_buf = torch.cuda.IntTensor([0])
|
||||
|
||||
def update(self, step: int, param: Tensor, grad: Tensor, exp_avg: Tensor, exp_avg_sq: Tensor):
|
||||
multi_tensor_applier(self.fused_adam, self.dummy_overflow_buf, [[grad], [param], [exp_avg], [exp_avg_sq]],
|
||||
self.lr, self.beta1, self.beta2, self.eps, step, self.use_adamw, True, self.weight_decay,
|
||||
-1)
|
||||
|
||||
|
||||
class CPUAdamKernel(AdamKernel):
|
||||
|
||||
def __init__(self, lr: float, beta1: float, beta2: float, eps: float, weight_decay: float, use_adamw: bool) -> None:
|
||||
super().__init__(lr, beta1, beta2, eps, weight_decay, use_adamw)
|
||||
from colossalai.kernel.op_builder import CPUAdamBuilder
|
||||
cpu_optim = CPUAdamBuilder().load()
|
||||
|
||||
self.cpu_adam_op = cpu_optim.CPUAdamOptimizer(lr, beta1, beta2, eps, weight_decay, use_adamw)
|
||||
|
||||
def update(self, step: int, param: Tensor, grad: Tensor, exp_avg: Tensor, exp_avg_sq: Tensor):
|
||||
self.cpu_adam_op.step(step, self.lr, self.beta1, self.beta2, self.eps, self.weight_decay, True, param.view(-1),
|
||||
grad.view(-1), exp_avg.view(-1), exp_avg_sq.view(-1), -1)
|
||||
|
||||
|
||||
def check_adam_kernel(kernel: Type[AdamKernel], adamw: bool, weight_decay: float, p_dtype: torch.dtype,
|
||||
g_dtype: torch.dtype, device: torch.device, n_steps: int, rtol: float, atol: float):
|
||||
lr = 1e-3
|
||||
beta1, beta2 = 0.9, 0.999
|
||||
eps = 1e-8
|
||||
torch_adam = TorchAdamKernel(lr, beta1, beta2, eps, weight_decay, adamw)
|
||||
adam_kernel = kernel(lr, beta1, beta2, eps, weight_decay, adamw)
|
||||
master_p = torch.rand(64, device=device)
|
||||
master_g = torch.rand_like(master_p)
|
||||
master_exp_avg = torch.zeros_like(master_p)
|
||||
master_exp_avg_sq = torch.zeros_like(master_p)
|
||||
p = master_p.clone().to(p_dtype)
|
||||
g = master_g.clone().to(g_dtype)
|
||||
exp_avg = master_exp_avg.clone()
|
||||
exp_avg_sq = master_exp_avg_sq.clone()
|
||||
|
||||
for step in range(1, 1 + n_steps):
|
||||
torch_adam.update(step, master_p, master_g, master_exp_avg, master_exp_avg_sq)
|
||||
adam_kernel.update(step, p, g, exp_avg, exp_avg_sq)
|
||||
# if overflow, the weight won't be updated. so there will be no nan in p
|
||||
assert not torch.isnan(p).any()
|
||||
assert torch.allclose(master_p, p.float(), rtol=rtol, atol=atol)
|
||||
|
||||
|
||||
@pytest.mark.parametrize('adamw', [False, True])
|
||||
@pytest.mark.parametrize('weight_decay', [0.0, 0.1])
|
||||
@pytest.mark.parametrize('p_dtype, g_dtype', _FUSED_ALLOWED_P_G_TYPES)
|
||||
def test_fused_adam_kernel(adamw, weight_decay, p_dtype, g_dtype):
|
||||
rtol, atol = 1e-5, 1e-8
|
||||
if p_dtype is torch.float16 or g_dtype is torch.float16:
|
||||
rtol, atol = 1e-3, 1e-3
|
||||
if p_dtype is torch.bfloat16 or g_dtype is torch.bfloat16:
|
||||
rtol, atol = 4e-3, 4e-3
|
||||
check_adam_kernel(FusedAdamKernel, adamw, weight_decay, p_dtype, g_dtype, get_current_device(), 3, rtol, atol)
|
||||
|
||||
|
||||
@pytest.mark.parametrize('adamw', [False, True])
|
||||
@pytest.mark.parametrize('weight_decay', [0.0, 0.1])
|
||||
@pytest.mark.parametrize('p_dtype, g_dtype', _CPU_ALLOWED_P_G_TYPES)
|
||||
def test_cpu_adam_kernel(adamw, weight_decay, p_dtype, g_dtype):
|
||||
rtol, atol = 1e-5, 1e-8
|
||||
if p_dtype is torch.float16 or g_dtype is torch.float16:
|
||||
rtol, atol = 1e-3, 1e-3
|
||||
check_adam_kernel(CPUAdamKernel, adamw, weight_decay, p_dtype, g_dtype, torch.device('cpu'), 3, rtol, atol)
|
86
tests/test_optimizer/test_adam_optim.py
Normal file
86
tests/test_optimizer/test_adam_optim.py
Normal file
@@ -0,0 +1,86 @@
|
||||
from copy import deepcopy
|
||||
from typing import Type, Union
|
||||
|
||||
import pytest
|
||||
import torch
|
||||
import torch.nn as nn
|
||||
from torch.optim import Adam, AdamW
|
||||
|
||||
from colossalai.nn.optimizer import CPUAdam, FusedAdam, HybridAdam
|
||||
from tests.kit.model_zoo import model_zoo
|
||||
|
||||
_ALLOWED_OPTIM_DEVICES = [
|
||||
(FusedAdam, torch.device('cuda:0')),
|
||||
(CPUAdam, torch.device('cpu')),
|
||||
(CPUAdam, torch.device('cuda:0')),
|
||||
(HybridAdam, torch.device('cpu')),
|
||||
(HybridAdam, torch.device('cuda:0')),
|
||||
]
|
||||
|
||||
_ALLOWED_P_G_TYPES = [
|
||||
(torch.float, torch.float), # pure fp32
|
||||
(torch.float, torch.half), # fp16 amp
|
||||
(torch.float, torch.bfloat16), # bfloat16 amp
|
||||
# (torch.half, torch.half), # FIXME(ver217): cpu adam kernel does not support pure fp16
|
||||
# (torch.bfloat16, torch.bfloat16), # FIXME(ver217): cpu adam kernel does not support pure bfloat16
|
||||
]
|
||||
|
||||
N_STEPS = 3
|
||||
|
||||
|
||||
def setup_param_groups(bert_model: nn.Module) -> list:
|
||||
no_decay = ["bias", "LayerNorm.weight"]
|
||||
optimizer_grouped_parameters = [
|
||||
{
|
||||
"params": [p for n, p in bert_model.named_parameters() if not any(nd in n for nd in no_decay)],
|
||||
"weight_decay": 0.1,
|
||||
},
|
||||
{
|
||||
"params": [p for n, p in bert_model.named_parameters() if any(nd in n for nd in no_decay)],
|
||||
"weight_decay": 0.0,
|
||||
},
|
||||
]
|
||||
return optimizer_grouped_parameters
|
||||
|
||||
|
||||
def set_grad(model: nn.Module, torch_model: nn.Module, g_dtype: torch.dtype) -> None:
|
||||
for p, torch_p in zip(model.parameters(), torch_model.parameters()):
|
||||
torch_p.grad = torch.rand_like(torch_p)
|
||||
# avoid inconsistent grad and param dtype error
|
||||
orig_p = p.data
|
||||
p.data = torch_p.grad.clone().to(g_dtype)
|
||||
p.grad = p.data
|
||||
p.data = orig_p
|
||||
|
||||
|
||||
@pytest.mark.parametrize('optim_cls, device', _ALLOWED_OPTIM_DEVICES)
|
||||
@pytest.mark.parametrize('adamw', [False, True])
|
||||
@pytest.mark.parametrize('p_dtype, g_dtype', _ALLOWED_P_G_TYPES)
|
||||
def test_adam_optim_on_bert(optim_cls: Union[Type[FusedAdam], Type[CPUAdam], Type[HybridAdam]], device: torch.device,
|
||||
adamw: bool, p_dtype: torch.dtype, g_dtype: torch.dtype) -> None:
|
||||
model_fn, *_ = next(iter(model_zoo.get_sub_registry('transformers_bert_for_sequence_classification').values()))
|
||||
torch_model = model_fn().to(device)
|
||||
model = deepcopy(torch_model).to(p_dtype)
|
||||
lr = 1e-3
|
||||
beta1, beta2 = 0.9, 0.999
|
||||
eps = 1e-8
|
||||
torch_optim_cls = AdamW if adamw else Adam
|
||||
torch_optim = torch_optim_cls(setup_param_groups(torch_model), lr=lr, betas=(beta1, beta2), eps=eps)
|
||||
optim = optim_cls(setup_param_groups(model), lr=lr, betas=(beta1, beta2), eps=eps, adamw_mode=adamw)
|
||||
|
||||
rtol, atol = 1e-5, 1e-5
|
||||
if p_dtype is torch.float16 or g_dtype is torch.float16:
|
||||
rtol, atol = 2e-3, 2e-3
|
||||
if p_dtype is torch.bfloat16 or g_dtype is torch.bfloat16:
|
||||
rtol, atol = 4e-3, 4e-3
|
||||
|
||||
for _ in range(N_STEPS):
|
||||
set_grad(model, torch_model, g_dtype)
|
||||
torch_optim.step()
|
||||
optim.step()
|
||||
torch_optim.zero_grad()
|
||||
optim.zero_grad()
|
||||
for p, torch_p in zip(model.parameters(), torch_model.parameters()):
|
||||
# if overflow, the weight won't be updated. so there will be no nan in p
|
||||
assert not torch.isnan(p).any()
|
||||
assert torch.allclose(p.float(), torch_p, rtol=rtol, atol=atol)
|
@@ -1,121 +0,0 @@
|
||||
import math
|
||||
|
||||
import torch
|
||||
|
||||
from colossalai.testing import clear_cache_before_run, parameterize
|
||||
|
||||
|
||||
def torch_adam_update(
|
||||
step,
|
||||
lr,
|
||||
beta1,
|
||||
beta2,
|
||||
eps,
|
||||
weight_decay,
|
||||
param,
|
||||
grad,
|
||||
exp_avg,
|
||||
exp_avg_sq,
|
||||
use_adamw,
|
||||
):
|
||||
bias_correction1 = 1 - beta1**step
|
||||
bias_correction2 = 1 - beta2**step
|
||||
|
||||
if weight_decay != 0:
|
||||
if use_adamw:
|
||||
# Perform stepweight decay
|
||||
param.mul_(1 - lr * weight_decay)
|
||||
else:
|
||||
grad = grad.add(param, alpha=weight_decay)
|
||||
|
||||
# Decay the first and second moment running average coefficient
|
||||
exp_avg.mul_(beta1).add_(grad, alpha=1 - beta1)
|
||||
exp_avg_sq.mul_(beta2).addcmul_(grad, grad, value=1 - beta2)
|
||||
denom = (exp_avg_sq.sqrt() / math.sqrt(bias_correction2)).add_(eps)
|
||||
|
||||
step_size = lr / bias_correction1
|
||||
|
||||
param.addcdiv_(exp_avg, denom, value=-step_size)
|
||||
|
||||
|
||||
def assertLess(data_diff, threshold, msg):
|
||||
assert data_diff < threshold, msg
|
||||
|
||||
|
||||
def assertTrue(condition, msg):
|
||||
assert condition, msg
|
||||
|
||||
|
||||
@clear_cache_before_run()
|
||||
@parameterize('adamw', [True, False])
|
||||
@parameterize('step', [1, 2])
|
||||
@parameterize('p_dtype', [torch.float, torch.half])
|
||||
@parameterize('g_dtype', [torch.float, torch.half])
|
||||
def test_cpu_adam(adamw, step, p_dtype, g_dtype):
|
||||
lr = 1e-3
|
||||
beta1, beta2 = 0.9, 0.999
|
||||
eps = 1e-8
|
||||
weight_decay = 0
|
||||
|
||||
for i in range(3):
|
||||
p_data = torch.rand(64, dtype=p_dtype)
|
||||
p_data_copy = p_data.clone().float()
|
||||
p_grad = torch.rand(64, dtype=g_dtype)
|
||||
p_grad_copy = p_grad.clone().float()
|
||||
exp_avg = torch.rand(p_data.shape)
|
||||
exp_avg_copy = exp_avg.clone()
|
||||
exp_avg_sq = torch.rand(p_data.shape)
|
||||
exp_avg_sq_copy = exp_avg_sq.clone()
|
||||
|
||||
from colossalai.kernel.op_builder import CPUAdamBuilder
|
||||
cpu_optim = CPUAdamBuilder().load()
|
||||
|
||||
cpu_adam_op = cpu_optim.CPUAdamOptimizer(lr, beta1, beta2, eps, weight_decay, adamw)
|
||||
|
||||
cpu_adam_op.step(
|
||||
step,
|
||||
lr,
|
||||
beta1,
|
||||
beta2,
|
||||
eps,
|
||||
weight_decay,
|
||||
True,
|
||||
p_data.view(-1), # fp32 data
|
||||
p_grad.view(-1), # fp32 grad
|
||||
exp_avg.view(-1),
|
||||
exp_avg_sq.view(-1),
|
||||
-1,
|
||||
)
|
||||
|
||||
torch_adam_update(
|
||||
step,
|
||||
lr,
|
||||
beta1,
|
||||
beta2,
|
||||
eps,
|
||||
weight_decay,
|
||||
p_data_copy, # fp32 data
|
||||
p_grad_copy, # fp32 grad
|
||||
exp_avg_copy,
|
||||
exp_avg_sq_copy,
|
||||
adamw,
|
||||
)
|
||||
var = p_data_copy - p_data
|
||||
data_diff = torch.max(torch.abs(var))
|
||||
threshold = 1e-3
|
||||
assertLess(
|
||||
data_diff,
|
||||
threshold,
|
||||
f"p_data diff {data_diff}. failed check, step {step}, lr {lr}, eps "
|
||||
f"{eps} beta1 {beta1} beta2 {beta2} weight_decay {weight_decay} p_dtype {p_dtype}, g_dtype {g_dtype}",
|
||||
)
|
||||
max_grad_diff = torch.max(torch.abs(p_grad_copy - p_grad))
|
||||
assertTrue(max_grad_diff < threshold, f"diff {max_grad_diff}")
|
||||
max_exp_avg_diff = torch.max(torch.abs(exp_avg_copy - exp_avg))
|
||||
assertTrue(max_exp_avg_diff < threshold, f"max_exp_avg_diff {max_exp_avg_diff}")
|
||||
max_exp_avg_sq_diff = torch.max(torch.abs(exp_avg_sq_copy - exp_avg_sq))
|
||||
assertTrue(max_exp_avg_sq_diff < threshold, f"max_exp_avg_sq_diff {max_exp_avg_sq_diff}")
|
||||
|
||||
|
||||
if __name__ == '__main__':
|
||||
test_cpu_adam()
|
@@ -1,64 +0,0 @@
|
||||
import torch
|
||||
import torch.nn as nn
|
||||
from torch.optim import AdamW
|
||||
from torch.optim.adam import Adam
|
||||
|
||||
from colossalai.nn.optimizer.fused_adam import FusedAdam
|
||||
from colossalai.testing import clear_cache_before_run, parameterize
|
||||
|
||||
|
||||
class FC(nn.Module):
|
||||
|
||||
def __init__(self) -> None:
|
||||
super().__init__()
|
||||
self.fc = nn.Sequential(nn.Linear(64, 64))
|
||||
|
||||
def forward(self, x):
|
||||
return self.fc(x)
|
||||
|
||||
|
||||
@clear_cache_before_run()
|
||||
@parameterize('adamw', [False, True])
|
||||
@parameterize('p_dtype', [torch.float, torch.half])
|
||||
@parameterize('g_dtype', [torch.float, torch.half])
|
||||
def test_adam(adamw, p_dtype, g_dtype):
|
||||
model = FC().cuda().to(p_dtype)
|
||||
state = model.state_dict()
|
||||
model_copy = FC().cuda().to(p_dtype)
|
||||
model_copy.load_state_dict(state.copy())
|
||||
|
||||
if adamw:
|
||||
optim = FusedAdam(model.parameters(), lr=1e-3, adamw_mode=True)
|
||||
torch_optim = AdamW(model_copy.parameters(), lr=1e-3)
|
||||
else:
|
||||
optim = FusedAdam(model.parameters(), lr=1e-3)
|
||||
torch_optim = Adam(model_copy.parameters(), lr=1e-3)
|
||||
|
||||
data = torch.rand(1024, 64).cuda().to(p_dtype)
|
||||
data_copy = data.clone()
|
||||
label = torch.rand(1024, 64).cuda().to(p_dtype)
|
||||
|
||||
for d, l in zip(data, label):
|
||||
y = model(d)
|
||||
loss = ((l - y)**2).sum()
|
||||
optim.zero_grad()
|
||||
loss.backward()
|
||||
if p_dtype != g_dtype:
|
||||
for i in range(len(optim.param_groups[0]['params'])):
|
||||
optim.param_groups[0]['params'][i].grad.data = optim.param_groups[0]['params'][i].grad.data.to(g_dtype)
|
||||
optim.step()
|
||||
|
||||
for d, l in zip(data_copy, label):
|
||||
y = model_copy(d)
|
||||
loss = ((l - y)**2).sum()
|
||||
torch_optim.zero_grad()
|
||||
loss.backward()
|
||||
torch_optim.step()
|
||||
|
||||
assert len(optim.param_groups[0]['params']) == len(torch_optim.param_groups[0]['params'])
|
||||
|
||||
for i in range(len(optim.param_groups[0]['params'])):
|
||||
if torch.isnan(optim.param_groups[0]['params'][i]).any() \
|
||||
or torch.isnan(torch_optim.param_groups[0]['params'][i]).any():
|
||||
continue
|
||||
assert torch.allclose(optim.param_groups[0]['params'][i], torch_optim.param_groups[0]['params'][i], 2e-3, 2e-3)
|
@@ -1,95 +0,0 @@
|
||||
import math
|
||||
|
||||
import torch
|
||||
import torch.nn as nn
|
||||
from numpy import dtype
|
||||
|
||||
from colossalai.testing import clear_cache_before_run, parameterize
|
||||
from colossalai.utils import multi_tensor_applier
|
||||
|
||||
|
||||
def torch_adam_update(
|
||||
step,
|
||||
lr,
|
||||
beta1,
|
||||
beta2,
|
||||
eps,
|
||||
weight_decay,
|
||||
param,
|
||||
grad,
|
||||
exp_avg,
|
||||
exp_avg_sq,
|
||||
use_adamw,
|
||||
):
|
||||
bias_correction1 = 1 - beta1**step
|
||||
bias_correction2 = 1 - beta2**step
|
||||
|
||||
if weight_decay != 0:
|
||||
if use_adamw:
|
||||
# Perform stepweight decay
|
||||
param.mul_(1 - lr * weight_decay)
|
||||
else:
|
||||
grad = grad.add(param, alpha=weight_decay)
|
||||
|
||||
# Decay the first and second moment running average coefficient
|
||||
exp_avg.mul_(beta1).add_(grad, alpha=1 - beta1)
|
||||
exp_avg_sq.mul_(beta2).addcmul_(grad, grad, value=1 - beta2)
|
||||
denom = (exp_avg_sq.sqrt() / math.sqrt(bias_correction2)).add_(eps)
|
||||
|
||||
step_size = lr / bias_correction1
|
||||
|
||||
param.addcdiv_(exp_avg, denom, value=-step_size)
|
||||
|
||||
|
||||
@clear_cache_before_run()
|
||||
@parameterize('adamw', [False, True])
|
||||
@parameterize('step', [1, 2])
|
||||
@parameterize('p_dtype', [torch.float, torch.half])
|
||||
@parameterize('g_dtype', [torch.float, torch.half])
|
||||
def test_adam(adamw, step, p_dtype, g_dtype):
|
||||
from colossalai.kernel.op_builder import FusedOptimBuilder
|
||||
fused_optim = FusedOptimBuilder().load()
|
||||
fused_adam = fused_optim.multi_tensor_adam
|
||||
|
||||
dummy_overflow_buf = torch.cuda.IntTensor([0])
|
||||
|
||||
count = 0
|
||||
|
||||
for i in range(3):
|
||||
p = torch.rand(64, dtype=p_dtype).cuda()
|
||||
p_copy = p.clone().float()
|
||||
g = torch.rand(p.shape, dtype=g_dtype).cuda()
|
||||
g_copy = g.clone().float()
|
||||
m = torch.rand(p.shape).cuda()
|
||||
m_copy = m.clone()
|
||||
v = torch.rand(p.shape).cuda()
|
||||
v_copy = v.clone()
|
||||
|
||||
lr = 1e-3
|
||||
beta1, beta2 = 0.9, 0.999
|
||||
eps = 1e-8
|
||||
weight_decay = 0
|
||||
|
||||
multi_tensor_applier(fused_adam, dummy_overflow_buf, [[g], [p], [m], [v]], lr, beta1, beta2, eps, step, adamw,
|
||||
True, weight_decay, -1)
|
||||
|
||||
torch_adam_update(
|
||||
step,
|
||||
lr,
|
||||
beta1,
|
||||
beta2,
|
||||
eps,
|
||||
weight_decay,
|
||||
p_copy, # fp32 data
|
||||
g_copy, # fp32 grad
|
||||
m_copy,
|
||||
v_copy,
|
||||
adamw,
|
||||
)
|
||||
|
||||
if torch.isnan(p).any() or torch.isnan(p_copy).any():
|
||||
count += 1
|
||||
continue
|
||||
assert count < 200, "too many nans"
|
||||
assert torch.allclose(p.to(torch.float), p_copy.to(torch.float), 1e-5,
|
||||
1e-5), f"failed check, adamw {adamw}, p_dtype {p_dtype}, g_dtype {g_dtype}"
|
@@ -1,42 +0,0 @@
|
||||
import torch
|
||||
import torch.nn as nn
|
||||
from torch.optim import AdamW
|
||||
from torch.optim.adam import Adam
|
||||
|
||||
from colossalai.nn.optimizer.hybrid_adam import HybridAdam
|
||||
from colossalai.testing import clear_cache_before_run, parameterize
|
||||
|
||||
RE = 3
|
||||
|
||||
|
||||
@clear_cache_before_run()
|
||||
@parameterize('adamw', [False, True])
|
||||
@parameterize('device', ['cpu', 'cuda:0'])
|
||||
@parameterize('p_dtype', [torch.float])
|
||||
@parameterize('g_dtype', [torch.float, torch.half])
|
||||
def test_adam(adamw, device, p_dtype, g_dtype):
|
||||
rng_state = torch.get_rng_state()
|
||||
p = nn.Parameter(torch.rand(64).to(device, p_dtype))
|
||||
torch.set_rng_state(rng_state)
|
||||
p_copy = nn.Parameter(torch.rand(64).to(device).float())
|
||||
|
||||
if adamw:
|
||||
optim = HybridAdam([p], lr=1e-3, adamw_mode=True)
|
||||
torch_optim = AdamW([p_copy], lr=1e-3)
|
||||
else:
|
||||
optim = HybridAdam([p], lr=1e-3)
|
||||
torch_optim = Adam([p_copy], lr=1e-3)
|
||||
|
||||
print(f"adaw mode {adamw}, device {device}, p_dtype {p_dtype}, g_dtype {g_dtype}")
|
||||
for i in range(RE):
|
||||
p.grad = torch.rand(64).to(device, p_dtype)
|
||||
p_copy.grad = p.grad.clone().float()
|
||||
p.grad.data = p.grad.data.to(g_dtype)
|
||||
|
||||
optim.step()
|
||||
torch_optim.step()
|
||||
|
||||
if torch.isnan(p.data).any() or torch.isnan(p_copy.data).any():
|
||||
continue
|
||||
assert torch.allclose(p.data, p_copy.data, 1e-4, 1e-2), \
|
||||
f"adaw mode {adamw}, device {device}, p_dtype {p_dtype}, g_dtype {g_dtype}"
|
@@ -21,23 +21,40 @@ TEST_MODELS = ['gpt2']
|
||||
# these models are too small, all parameters in these models are compacted into one chunk
|
||||
EXAMPLE_MODELS = ['albert', 'beit', 'bert', 'hanging_param_model', 'nested_model', 'repeated_computed_layers']
|
||||
|
||||
# bfloat16 cannot represent them exactly
|
||||
BF16_IGNORED_KEYS = [
|
||||
'albert.embeddings.word_embeddings.weight',
|
||||
'albert.embeddings.position_embeddings.weight',
|
||||
'masked_bias',
|
||||
]
|
||||
|
||||
def check_param(model: ZeroDDP, torch_model: torch.nn.Module):
|
||||
zero_dict = model.state_dict(only_rank_0=False)
|
||||
|
||||
def check_param(model: ZeroDDP, torch_model: torch.nn.Module, dtype: torch.dtype):
|
||||
zero_dict = model.state_dict(only_rank_0=False, dtype=dtype)
|
||||
torch_dict = torch_model.state_dict()
|
||||
|
||||
for key, value in torch_dict.items():
|
||||
# key is 'module.model.PARAMETER', so we truncate it
|
||||
key = key[7:]
|
||||
assert key in zero_dict, "{} not in ZeRO dictionary.".format(key)
|
||||
temp_zero_value = zero_dict[key].to(device=value.device, dtype=value.dtype)
|
||||
temp_zero_value = zero_dict[key].to(device=value.device)
|
||||
if dtype is torch.bfloat16 and any(k in key for k in BF16_IGNORED_KEYS):
|
||||
continue
|
||||
rtol, atol = 1e-3, 4e-3
|
||||
if dtype is torch.bfloat16:
|
||||
rtol, atol = 4e-3, 8e-3
|
||||
# debug_print([0], "max range: ", key, torch.max(torch.abs(value - temp_zero_value)))
|
||||
assert_close(value, temp_zero_value, rtol=1e-3, atol=4e-3)
|
||||
assert_close(value.float(),
|
||||
temp_zero_value.float(),
|
||||
rtol=rtol,
|
||||
atol=atol,
|
||||
msg=lambda s: s + f'\n{key}\n{temp_zero_value.dtype}')
|
||||
|
||||
|
||||
@parameterize('placement_policy', ['cuda', 'cpu', 'auto', 'const'])
|
||||
@parameterize('model_name', TEST_MODELS)
|
||||
def exam_model_step(placement_policy, model_name: str):
|
||||
@parameterize('mixed_precision', [torch.half, torch.bfloat16])
|
||||
def exam_model_step(placement_policy, model_name: str, mixed_precision: torch.dtype):
|
||||
set_seed(42)
|
||||
get_components_func = non_distributed_component_funcs.get_callable(model_name)
|
||||
model_builder, train_dataloader, test_dataloader, optimizer_class, criterion = get_components_func()
|
||||
@@ -65,7 +82,7 @@ def exam_model_step(placement_policy, model_name: str):
|
||||
init_device = None
|
||||
chunk_manager = ChunkManager(config_dict, init_device=init_device)
|
||||
gemini_manager = GeminiManager(placement_policy, chunk_manager)
|
||||
model = ZeroDDP(model, gemini_manager, pin_memory=True)
|
||||
model = ZeroDDP(model, gemini_manager, pin_memory=True, mixed_precision=mixed_precision)
|
||||
|
||||
optimizer = HybridAdam(model.parameters(), lr=1e-3)
|
||||
zero_optim = ZeroOptimizer(optimizer, model, initial_scale=128)
|
||||
@@ -74,6 +91,7 @@ def exam_model_step(placement_policy, model_name: str):
|
||||
torch_model.eval()
|
||||
|
||||
set_seed(dist.get_rank() * 3 + 128)
|
||||
rtol, atol = 1e-4, 1e-5
|
||||
for i, (input_ids, label) in enumerate(train_dataloader):
|
||||
if i > 2:
|
||||
break
|
||||
@@ -83,17 +101,18 @@ def exam_model_step(placement_policy, model_name: str):
|
||||
|
||||
torch_loss = run_fwd_bwd(torch_model, input_ids, label, criterion, torch_optim)
|
||||
loss = run_fwd_bwd(model, input_ids, label, criterion, zero_optim)
|
||||
assert_close(torch_loss, loss)
|
||||
assert_close(torch_loss, loss, rtol=rtol, atol=atol)
|
||||
|
||||
zero_optim.step()
|
||||
torch_optim.step()
|
||||
|
||||
check_param(model, torch_model)
|
||||
check_param(model, torch_model, mixed_precision)
|
||||
|
||||
|
||||
@parameterize('placement_policy', ['cuda', 'cpu', 'auto', 'const'])
|
||||
@parameterize('model_name', EXAMPLE_MODELS)
|
||||
def exam_tiny_example(placement_policy, model_name: str):
|
||||
@parameterize('mixed_precision', [torch.half, torch.bfloat16])
|
||||
def exam_tiny_example(placement_policy, model_name: str, mixed_precision: torch.dtype):
|
||||
set_seed(2008)
|
||||
get_components_func = non_distributed_component_funcs.get_callable(model_name)
|
||||
model_builder, train_dataloader, test_dataloader, optimizer_class, criterion = get_components_func()
|
||||
@@ -113,7 +132,7 @@ def exam_tiny_example(placement_policy, model_name: str):
|
||||
|
||||
chunk_manager = init_chunk_manager(model=model, init_device=get_current_device(), search_range_mb=1)
|
||||
gemini_manager = GeminiManager(placement_policy, chunk_manager)
|
||||
model = ZeroDDP(model, gemini_manager, pin_memory=True)
|
||||
model = ZeroDDP(model, gemini_manager, pin_memory=True, mixed_precision=mixed_precision)
|
||||
optimizer = HybridAdam(model.parameters(), lr=1e-3)
|
||||
zero_optim = ZeroOptimizer(optimizer, model, initial_scale=2)
|
||||
|
||||
@@ -121,6 +140,9 @@ def exam_tiny_example(placement_policy, model_name: str):
|
||||
torch_model.eval()
|
||||
|
||||
set_seed(dist.get_rank() * 3 + 128)
|
||||
rtol, atol = 1.5e-6, 2e-5
|
||||
if mixed_precision is torch.bfloat16:
|
||||
rtol, atol = 2e-3, 2e-3
|
||||
for i, (input_ids, label) in enumerate(train_dataloader):
|
||||
if i > 2:
|
||||
break
|
||||
@@ -133,12 +155,12 @@ def exam_tiny_example(placement_policy, model_name: str):
|
||||
|
||||
torch_loss = run_fwd_bwd(torch_model, input_ids, label, criterion, torch_optim)
|
||||
loss = run_fwd_bwd(model, input_ids, label, criterion, zero_optim)
|
||||
assert_close(torch_loss, loss, rtol=1.5e-6, atol=2e-5) # atol should be 2e-5 for torch lower than 1.12
|
||||
assert_close(torch_loss, loss, rtol=rtol, atol=atol) # atol should be 2e-5 for torch lower than 1.12
|
||||
|
||||
zero_optim.step()
|
||||
torch_optim.step()
|
||||
|
||||
check_param(model, torch_model)
|
||||
check_param(model, torch_model, mixed_precision)
|
||||
|
||||
|
||||
def run_dist(rank, world_size, port):
|
||||
|
@@ -16,7 +16,11 @@ from colossalai.zero.low_level._utils import has_inf_or_nan
|
||||
from tests.components_to_test.registry import non_distributed_component_funcs
|
||||
|
||||
|
||||
def run_dist(rank, world_size, port, parallel_config):
|
||||
def run_dist(rank, world_size, port, parallel_config, bf16):
|
||||
is_mp_config = parallel_config == MP_PARALLEL_CONFIG
|
||||
is_zero_config = parallel_config == ZERO_PARALLEL_CONFIG
|
||||
if bf16:
|
||||
parallel_config['zero']['model_config']['bf16'] = True
|
||||
colossalai.launch(config=parallel_config,
|
||||
rank=rank,
|
||||
world_size=world_size,
|
||||
@@ -30,7 +34,8 @@ def run_dist(rank, world_size, port, parallel_config):
|
||||
model_builder, train_dataloader, _, optimizer_class, criterion = get_components_func()
|
||||
with ZeroInitContext(target_device=torch.cuda.current_device(),
|
||||
shard_strategy=gpc.config.zero.model_config.shard_strategy,
|
||||
shard_param=True):
|
||||
shard_param=True,
|
||||
bf16=bf16):
|
||||
colo_model = model_builder(checkpoint=True)
|
||||
|
||||
colo_optimizer = optimizer_class(colo_model.parameters(), lr=1e-3)
|
||||
@@ -38,7 +43,8 @@ def run_dist(rank, world_size, port, parallel_config):
|
||||
optimizer=colo_optimizer,
|
||||
criterion=criterion,
|
||||
train_dataloader=train_dataloader)
|
||||
torch_model = model_builder(checkpoint=True).half()
|
||||
dtype = torch.bfloat16 if bf16 else torch.float16
|
||||
torch_model = model_builder(checkpoint=True).to(dtype)
|
||||
col_model_deepcopy(engine.model, torch_model)
|
||||
torch_model = torch_model.cuda().float()
|
||||
|
||||
@@ -80,9 +86,9 @@ def run_dist(rank, world_size, port, parallel_config):
|
||||
torch_optimizer.step()
|
||||
i += 1
|
||||
|
||||
if parallel_config == MP_PARALLEL_CONFIG:
|
||||
if is_mp_config:
|
||||
check_params(torch_model, colo_model, loose=True)
|
||||
elif parallel_config == ZERO_PARALLEL_CONFIG:
|
||||
elif is_zero_config:
|
||||
check_sharded_model_params(torch_model, colo_model, loose=True)
|
||||
|
||||
|
||||
@@ -97,9 +103,10 @@ def test_mp_engine(world_size):
|
||||
|
||||
@pytest.mark.dist
|
||||
@pytest.mark.parametrize("world_size", [1, 2])
|
||||
@pytest.mark.parametrize("bf16", [True, False])
|
||||
@rerun_if_address_is_in_use()
|
||||
def test_zero_engine(world_size):
|
||||
spawn(run_dist, world_size, parallel_config=ZERO_PARALLEL_CONFIG)
|
||||
def test_zero_engine(world_size, bf16):
|
||||
spawn(run_dist, world_size, parallel_config=ZERO_PARALLEL_CONFIG, bf16=bf16)
|
||||
|
||||
|
||||
if __name__ == '__main__':
|
||||
|
@@ -82,7 +82,6 @@ def exam_zero_1_2_grad_acc():
|
||||
|
||||
def exam_zero_1_grad_acc():
|
||||
local_rank = torch.distributed.get_rank()
|
||||
grad_scale = 32
|
||||
seed_all(2008)
|
||||
|
||||
# create models
|
||||
@@ -101,7 +100,6 @@ def exam_zero_1_grad_acc():
|
||||
# level 1 and 2 will produce exactly the same results
|
||||
zero_optimizer = LowLevelZeroOptimizer(zero_optimizer,
|
||||
overlap_communication=False,
|
||||
initial_scale=grad_scale,
|
||||
reduce_bucket_size=262144,
|
||||
clip_grad_norm=1.0)
|
||||
|
||||
@@ -128,9 +126,8 @@ def exam_zero_1_grad_acc():
|
||||
if check_flag:
|
||||
# check grad
|
||||
for (n, p), z1p in zip(torch_model.named_parameters(), zero_model.parameters()):
|
||||
unscale_grad = z1p.grad / grad_scale
|
||||
# print(n, p.shape, torch.max(torch.abs(p.grad - unscale_grad)))
|
||||
assert torch.equal(p.grad, unscale_grad)
|
||||
assert torch.equal(p.grad, z1p.grad)
|
||||
|
||||
zero_optimizer._sync_grad()
|
||||
|
||||
|
@@ -7,7 +7,7 @@ from torch.nn.parallel import DistributedDataParallel as DDP
|
||||
from torch.testing import assert_close
|
||||
|
||||
import colossalai
|
||||
from colossalai.testing import rerun_if_address_is_in_use, spawn
|
||||
from colossalai.testing import parameterize, rerun_if_address_is_in_use, spawn
|
||||
from colossalai.testing.random import seed_all
|
||||
from colossalai.zero import LowLevelZeroOptimizer
|
||||
|
||||
@@ -25,15 +25,18 @@ class MlpModel(nn.Module):
|
||||
return x
|
||||
|
||||
|
||||
def half_close(a, b, loose=False):
|
||||
def loose_close(a, b, dtype: torch.dtype = torch.float32):
|
||||
rtol = None
|
||||
atol = None
|
||||
if loose:
|
||||
if dtype is torch.float16:
|
||||
rtol = 5e-2
|
||||
atol = 5e-4
|
||||
elif dtype is torch.bfloat16:
|
||||
rtol = 4e-3
|
||||
atol = 4e-3
|
||||
|
||||
a = a.detach().half()
|
||||
b = b.detach().half()
|
||||
a = a.detach().to(dtype)
|
||||
b = b.detach().to(dtype)
|
||||
|
||||
assert_close(a, b, rtol=rtol, atol=atol)
|
||||
|
||||
@@ -96,7 +99,8 @@ def exam_zero_1_2():
|
||||
assert torch.equal(z1p.data, z2p.data)
|
||||
|
||||
|
||||
def exam_zero_1_torch_ddp():
|
||||
@parameterize('dtype', [torch.float16, torch.bfloat16])
|
||||
def exam_zero_1_torch_ddp(dtype: torch.dtype):
|
||||
"""
|
||||
In this test, two pairs of model and optimizers are created.
|
||||
1. zero: use sharded optimizer and fp16 parameters
|
||||
@@ -109,15 +113,10 @@ def exam_zero_1_torch_ddp():
|
||||
seed_all(1453)
|
||||
|
||||
# create models
|
||||
zero_model = MlpModel()
|
||||
torch_model = copy.deepcopy(zero_model)
|
||||
torch_model = MlpModel().cuda()
|
||||
zero_model = copy.deepcopy(torch_model).to(dtype)
|
||||
|
||||
zero_model = zero_model.cuda().half()
|
||||
torch_model = DDP(torch_model.cuda(), bucket_cap_mb=0)
|
||||
torch_model = torch_model.cuda()
|
||||
|
||||
# for (n, p), z1p in zip(torch_model.named_parameters(), zero_model.parameters()):
|
||||
# half_close(p.data, z1p.data)
|
||||
torch_model = DDP(torch_model.cuda(), bucket_cap_mb=0).cuda()
|
||||
|
||||
# create optimizer
|
||||
zero_optimizer = torch.optim.SGD(zero_model.parameters(), lr=1)
|
||||
@@ -137,11 +136,11 @@ def exam_zero_1_torch_ddp():
|
||||
input_data = torch.rand(32, 128).cuda()
|
||||
|
||||
# zero-dp forward
|
||||
zero_output = zero_model(input_data.half())
|
||||
zero_output = zero_model(input_data.to(dtype))
|
||||
|
||||
# torch-ddp forward
|
||||
torch_output = torch_model(input_data)
|
||||
half_close(zero_output, torch_output, loose=True)
|
||||
loose_close(zero_output, torch_output, dtype=dtype)
|
||||
|
||||
# zero-dp backward
|
||||
zero_optimizer.backward(zero_output.mean().float(), sync_grad=False)
|
||||
@@ -151,7 +150,7 @@ def exam_zero_1_torch_ddp():
|
||||
|
||||
# check grad
|
||||
for (n, p), z1p in zip(torch_model.named_parameters(), zero_model.parameters()):
|
||||
half_close(p.grad, z1p.grad, loose=True)
|
||||
loose_close(p.grad, z1p.grad, dtype=dtype)
|
||||
|
||||
# zero-dp step
|
||||
zero_optimizer._sync_grad()
|
||||
@@ -163,7 +162,7 @@ def exam_zero_1_torch_ddp():
|
||||
# check updated param
|
||||
for (n, p), z1p in zip(torch_model.named_parameters(), zero_model.parameters()):
|
||||
# print(n, torch.max(torch.abs(p.data - z1p.data)))
|
||||
half_close(p.data, z1p.data, loose=True)
|
||||
loose_close(p.data, z1p.data, dtype=dtype)
|
||||
|
||||
|
||||
def run_dist(rank, world_size, port):
|
||||
|
Reference in New Issue
Block a user