Skip to content

Commit 6ba5521

Browse files
tchatonSeanNaren
authored andcommitted
[bug-fix] DDP and automatic_optimization=False (#4485)
* resolve bug * add self._running_manual_optim * update * update tests * update lightning module * resolve bug * update tests * update * resolve pep8 * update * replace by `ddp_spawn` * temporary fix * update * update * move update to training_loop * make both ddp_spawn * introduce `manual_optimizer_step` * update changelog * added changelog wrong place * add force_optimizer_step * update docstring for tests * update optimizer_step * update zero_grad * resolve flake8 * move update into manual_optimizer_step * add zero_grad * remove zero_grad tests * remove manual_backward in AMP, it doesn't help * update * loosen tests * update * update doc * add TODO * Removed unnecessary get model from native amp * Remove try except with pytest raise * Add seed, clean up imports, remove try catch to reproduce error * update code * update test * revert back * formatting * Update pytorch_lightning/core/lightning.py Co-authored-by: Jirka Borovec <[email protected]> Co-authored-by: SeanNaren <[email protected]> Co-authored-by: Sean Naren <[email protected]> Co-authored-by: Jirka Borovec <[email protected]> (cherry picked from commit 7e08b0d)
1 parent 760b83d commit 6ba5521

File tree

8 files changed

+393
-23
lines changed

8 files changed

+393
-23
lines changed

.gitignore

Lines changed: 1 addition & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -33,6 +33,7 @@ timit_data/
3333
.Python
3434
ide_layouts/
3535
build/
36+
_build/
3637
develop-eggs/
3738
dist/
3839
downloads/

docs/source/lightning_module.rst

Lines changed: 6 additions & 0 deletions
Original file line numberDiff line numberDiff line change
@@ -1009,6 +1009,12 @@ manual_backward
10091009
.. automethod:: pytorch_lightning.core.lightning.LightningModule.manual_backward
10101010
:noindex:
10111011

1012+
manual_optimizer_step
1013+
~~~~~~~~~~~~~~~~~~~~~
1014+
1015+
.. automethod:: pytorch_lightning.core.lightning.LightningModule.manual_optimizer_step
1016+
:noindex:
1017+
10121018
on_after_backward
10131019
~~~~~~~~~~~~~~~~~
10141020

docs/source/optimizers.rst

Lines changed: 3 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -36,17 +36,16 @@ to manually manage the optimization process. To do so, do the following:
3636
3737
# use self.backward which will also handle scaling the loss when using amp
3838
self.manual_backward(loss_a, opt_g)
39-
opt_g.step()
40-
opt_g.zero_grad()
39+
self.manual_optimizer_step(opt_g)
40+
4141
4242
# do anything you want
4343
loss_b = ...
4444
4545
# pass in any args that loss.backward() normally takes
4646
self.manual_backward(loss_b, opt_d, retain_graph=True)
4747
self.manual_backward(loss_b, opt_d)
48-
opt_d.step()
49-
opt_d.zero_grad()
48+
self.manual_optimizer_step(opt_d)
5049
5150
# log losses
5251
self.log('loss_a', loss_a)

pytorch_lightning/accelerators/accelerator.py

Lines changed: 5 additions & 4 deletions
Original file line numberDiff line numberDiff line change
@@ -109,10 +109,11 @@ def backward(self, closure_loss, optimizer, opt_idx, *args, **kwargs):
109109
def optimizer_step(self, optimizer, batch_idx, opt_idx, lambda_closure):
110110
model_ref = self.trainer.get_model()
111111
is_lbfgs = isinstance(optimizer, torch.optim.LBFGS)
112-
native_amp = self.trainer.amp_backend == AMPType.NATIVE
112+
using_native_amp = self.trainer.amp_backend == AMPType.NATIVE
113+
automatic_optimization = self.trainer.train_loop.automatic_optimization
113114

114115
# native amp + lbfgs is a no go right now
115-
if native_amp and is_lbfgs:
116+
if using_native_amp and is_lbfgs:
116117
raise MisconfigurationException(
117118
'native PyTorch amp and lbfgs are not compatible.'
118119
' To request, please file a Github issue in PyTorch and tag @mcarilli')
@@ -125,12 +126,12 @@ def optimizer_step(self, optimizer, batch_idx, opt_idx, lambda_closure):
125126
optimizer_idx=opt_idx,
126127
optimizer_closure=lambda_closure,
127128
on_tpu=False, # TPUAccelerator class sets this as True
128-
using_native_amp=native_amp,
129+
using_native_amp=using_native_amp,
129130
using_lbfgs=is_lbfgs
130131
)
131132

132133
# scale when native amp
133-
if native_amp:
134+
if automatic_optimization and using_native_amp:
134135
self.trainer.scaler.update()
135136

136137
def optimizer_zero_grad(self, batch_idx, optimizer, opt_idx):

pytorch_lightning/core/lightning.py

Lines changed: 49 additions & 1 deletion
Original file line numberDiff line numberDiff line change
@@ -111,6 +111,7 @@ def __init__(self, *args, **kwargs):
111111
self._datamodule = None
112112
self._results: Optional[Result] = None
113113
self._current_fx_name = ''
114+
self._running_manual_backward = False
114115

115116
def optimizers(self):
116117
opts = self.trainer.optimizers
@@ -1070,19 +1071,65 @@ def manual_backward(self, loss: Tensor, optimizer: Optimizer, *args, **kwargs) -
10701071
10711072
.. tip:: In manual mode we still automatically clip grads if Trainer(gradient_clip_val=x) is set
10721073
1074+
.. tip:: In manual mode we still automatically accumulate grad over batches if Trainer(accumulate_grad_batches=x) is set
1075+
and you use `model.manual_optimizer_step(optimizer)`
1076+
10731077
Example::
10741078
10751079
def training_step(...):
10761080
(opt_a, opt_b) = self.optimizers()
10771081
loss = ...
10781082
# automatically applies scaling, etc...
10791083
self.manual_backward(loss, opt_a)
1084+
self.manual_optimizer_step(opt_a)
10801085
"""
10811086
# make sure we're using manual opt
10821087
self._verify_is_manual_optimization('manual_backward')
10831088

10841089
# backward
1090+
self._running_manual_backward = True
10851091
self.trainer.train_loop.backward(loss, optimizer, -1, *args, **kwargs)
1092+
self._running_manual_backward = False
1093+
1094+
def manual_optimizer_step(self, optimizer: Optimizer, force_optimizer_step:bool = False) -> None:
1095+
"""
1096+
Call this directly from your training_step when doing optimizations manually.
1097+
By using this we can ensure that all the proper scaling when using 16-bit etc has been done for you
1098+
1099+
.. tip:: In manual mode we still automatically accumulate grad over batches if Trainer(accumulate_grad_batches=x) is set.
1100+
1101+
Args:
1102+
optimizer: Optimizer used to perform `.step()` call
1103+
1104+
force_optimizer_step: Whether to force an optimizer step. Could be useful when having 2 optimizers
1105+
and one should use accumulated gradients but not the other one.
1106+
One could put its own logic to force an optimizer step.
1107+
1108+
Example::
1109+
1110+
def training_step(...):
1111+
(opt_a, opt_b) = self.optimizers()
1112+
loss = ...
1113+
# automatically applies scaling, etc...
1114+
self.manual_backward(loss, opt_a)
1115+
# This will force an opt.step() even if accumulate_grad_batches is set.
1116+
self.manual_optimizer_step(opt_a, force_optimizer_step=True)
1117+
1118+
"""
1119+
# make sure we're using manual opt
1120+
self._verify_is_manual_optimization('manual_optimizer_step')
1121+
1122+
if not self.trainer.train_loop.should_accumulate() or force_optimizer_step:
1123+
1124+
# mock closure function as the user is responsible to call `manual_backward`
1125+
def mock_optimizer_closure():
1126+
return
1127+
1128+
self.trainer.train_loop.optimizer_step(optimizer, None, self.trainer.batch_idx, mock_optimizer_closure)
1129+
1130+
# update will be called after every optimizer_step call
1131+
if self.trainer.amp_backend == AMPType.NATIVE:
1132+
self.trainer.scaler.update()
10861133

10871134
def backward(self, loss: Tensor, optimizer: Optimizer, optimizer_idx: int, *args, **kwargs) -> None:
10881135
"""
@@ -1103,7 +1150,8 @@ def backward(self, loss, optimizer, optimizer_idx):
11031150
loss.backward()
11041151
11051152
"""
1106-
loss.backward(*args, **kwargs)
1153+
if self.trainer.train_loop.automatic_optimization or self._running_manual_backward:
1154+
loss.backward(*args, **kwargs)
11071155

11081156
def toggle_optimizer(self, optimizer: Optimizer, optimizer_idx: int):
11091157
"""

pytorch_lightning/trainer/training_loop.py

Lines changed: 53 additions & 5 deletions
Original file line numberDiff line numberDiff line change
@@ -303,6 +303,12 @@ def on_after_backward(self, training_step_output, batch_idx, untouched_loss):
303303
# when in dev debugging track the losses
304304
self.trainer.dev_debugger.track_train_loss_history(batch_idx, untouched_loss.detach())
305305

306+
def _check_training_step_output(self, training_step_output):
307+
if isinstance(training_step_output, torch.Tensor) and not self.automatic_optimization:
308+
if training_step_output.grad_fn is None:
309+
# TODO: Find why - RuntimeError: Expected to mark a variable ready only once ...
310+
raise MisconfigurationException("In manual optimization, `training_step` should not return a Tensor")
311+
306312
def training_step(self, split_batch, batch_idx, opt_idx, hiddens):
307313
# give the PL module a result for logging
308314
model = self.trainer.get_model()
@@ -312,6 +318,8 @@ def training_step(self, split_batch, batch_idx, opt_idx, hiddens):
312318
with self.trainer.profiler.profile("model_forward"):
313319
args = self.build_train_args(split_batch, batch_idx, opt_idx, hiddens)
314320
training_step_output = self.trainer.accelerator_backend.training_step(args)
321+
self._check_training_step_output(training_step_output)
322+
315323
training_step_output = self.trainer.call_hook("training_step_end", training_step_output)
316324

317325
training_step_output_for_epoch_end, training_step_output = self._process_training_step_output(
@@ -723,6 +731,8 @@ def train_step_and_backward_closure():
723731

724732
if self._curr_step_result is None:
725733
# user decided to skip optimization
734+
# make sure to zero grad.
735+
self.zero_grad_handler(batch_idx, optimizer, opt_idx)
726736
continue
727737

728738
batch_outputs = self._process_closure_result(
@@ -735,11 +745,8 @@ def train_step_and_backward_closure():
735745
grad_norm_dic = self._cur_grad_norm_dict
736746
self._cur_grad_norm_dict = None
737747

738-
# hook
739-
self.on_before_zero_grad(optimizer)
740-
741-
# clear gradients
742-
self.optimizer_zero_grad(batch_idx, optimizer, opt_idx)
748+
# hook + clear gradients
749+
self.zero_grad_handler(batch_idx, optimizer, opt_idx)
743750

744751
accumulated_loss = self.accumulated_loss.mean()
745752

@@ -949,3 +956,44 @@ def process_train_step_outputs(self, all_train_step_outputs, early_stopping_accu
949956
epoch_end_outputs.append(optimizer_idx_outputs)
950957

951958
return epoch_end_outputs
959+
960+
def prepare_optimizers(self):
961+
# in manual optimization we loop over all optimizers at once
962+
optimizers = self.get_optimizers_iterable()
963+
if not self.automatic_optimization:
964+
optimizers = [optimizers[0]]
965+
return optimizers
966+
967+
def run_train_split_start(self, split_idx, split_batch, opt_idx, optimizer):
968+
# set split_idx to trainer for tracking
969+
self.trainer.split_idx = split_idx
970+
971+
# make sure only the gradients of the current optimizer's parameters are calculated
972+
# in the training step to prevent dangling gradients in multiple-optimizer setup.
973+
if self.automatic_optimization and len(self.trainer.optimizers) > 1:
974+
model = self.trainer.get_model()
975+
model.toggle_optimizer(optimizer, opt_idx)
976+
977+
# use to track metrics internally
978+
self.trainer.logger_connector.on_train_split_start(split_idx, opt_idx, split_batch)
979+
980+
def update_running_loss(self):
981+
accumulated_loss = self.accumulated_loss.mean()
982+
983+
if accumulated_loss is not None:
984+
# calculate running loss for display
985+
self.running_loss.append(self.accumulated_loss.mean() * self.trainer.accumulate_grad_batches)
986+
987+
# reset for next set of accumulated grads
988+
self.accumulated_loss.reset()
989+
990+
def zero_grad_handler(self, batch_idx, optimizer, opt_idx):
991+
if self.automatic_optimization:
992+
# hook
993+
self.on_before_zero_grad(optimizer)
994+
optimizers = enumerate([optimizer])
995+
else:
996+
optimizers = self.get_optimizers_iterable()
997+
998+
for idx, optimizer in optimizers:
999+
self.optimizer_zero_grad(batch_idx, optimizer, opt_idx)

0 commit comments

Comments
 (0)