Update Trainer property docstrings (#16989)
This commit is contained in:
parent
20374b93f4
commit
279e3add00
|
@ -1150,6 +1150,14 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def log_dir(self) -> Optional[str]:
|
def log_dir(self) -> Optional[str]:
|
||||||
|
"""The directory for the current experiment. Use this to save images to, etc...
|
||||||
|
|
||||||
|
.. code-block:: python
|
||||||
|
|
||||||
|
def training_step(self, batch, batch_idx):
|
||||||
|
img = ...
|
||||||
|
save_img(img, self.trainer.log_dir)
|
||||||
|
"""
|
||||||
if len(self.loggers) > 0:
|
if len(self.loggers) > 0:
|
||||||
if not isinstance(self.loggers[0], TensorBoardLogger):
|
if not isinstance(self.loggers[0], TensorBoardLogger):
|
||||||
dirpath = self.loggers[0].save_dir
|
dirpath = self.loggers[0].save_dir
|
||||||
|
@ -1163,6 +1171,14 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def is_global_zero(self) -> bool:
|
def is_global_zero(self) -> bool:
|
||||||
|
"""Whether this process is the global zero in multi-node training.
|
||||||
|
|
||||||
|
.. code-block:: python
|
||||||
|
|
||||||
|
def training_step(self, batch, batch_idx):
|
||||||
|
if self.trainer.is_global_zero:
|
||||||
|
print("in node 0, accelerator 0")
|
||||||
|
"""
|
||||||
return self.strategy.is_global_zero
|
return self.strategy.is_global_zero
|
||||||
|
|
||||||
@property
|
@property
|
||||||
|
@ -1236,7 +1252,7 @@ class Trainer:
|
||||||
def ckpt_path(self, ckpt_path: Optional[_PATH]) -> None:
|
def ckpt_path(self, ckpt_path: Optional[_PATH]) -> None:
|
||||||
"""Allows you to manage which checkpoint is loaded statefully.
|
"""Allows you to manage which checkpoint is loaded statefully.
|
||||||
|
|
||||||
Examples::
|
.. code-block:: python
|
||||||
|
|
||||||
trainer = Trainer()
|
trainer = Trainer()
|
||||||
trainer.ckpt_path = "my/checkpoint/file.ckpt"
|
trainer.ckpt_path = "my/checkpoint/file.ckpt"
|
||||||
|
@ -1384,11 +1400,13 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def train_dataloader(self) -> TRAIN_DATALOADERS:
|
def train_dataloader(self) -> TRAIN_DATALOADERS:
|
||||||
|
"""The training dataloader(s) used during ``trainer.fit()``."""
|
||||||
if (combined_loader := self.fit_loop._combined_loader) is not None:
|
if (combined_loader := self.fit_loop._combined_loader) is not None:
|
||||||
return combined_loader.iterables
|
return combined_loader.iterables
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def val_dataloaders(self) -> EVAL_DATALOADERS:
|
def val_dataloaders(self) -> EVAL_DATALOADERS:
|
||||||
|
"""The validation dataloader(s) used during ``trainer.fit()`` or ``trainer.validate()``."""
|
||||||
if (combined_loader := self.fit_loop.epoch_loop.val_loop._combined_loader) is not None:
|
if (combined_loader := self.fit_loop.epoch_loop.val_loop._combined_loader) is not None:
|
||||||
return combined_loader.iterables
|
return combined_loader.iterables
|
||||||
elif (combined_loader := self.validate_loop._combined_loader) is not None:
|
elif (combined_loader := self.validate_loop._combined_loader) is not None:
|
||||||
|
@ -1396,25 +1414,32 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def test_dataloaders(self) -> EVAL_DATALOADERS:
|
def test_dataloaders(self) -> EVAL_DATALOADERS:
|
||||||
|
"""The test dataloader(s) used during ``trainer.test()``."""
|
||||||
if (combined_loader := self.test_loop._combined_loader) is not None:
|
if (combined_loader := self.test_loop._combined_loader) is not None:
|
||||||
return combined_loader.iterables
|
return combined_loader.iterables
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def predict_dataloaders(self) -> EVAL_DATALOADERS:
|
def predict_dataloaders(self) -> EVAL_DATALOADERS:
|
||||||
|
"""The prediction dataloader(s) used during ``trainer.predict()``."""
|
||||||
if (combined_loader := self.predict_loop._combined_loader) is not None:
|
if (combined_loader := self.predict_loop._combined_loader) is not None:
|
||||||
return combined_loader.iterables
|
return combined_loader.iterables
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def num_training_batches(self) -> Union[int, float]:
|
def num_training_batches(self) -> Union[int, float]:
|
||||||
|
"""The number of training batches that will be used during ``trainer.fit()``."""
|
||||||
return self.fit_loop.max_batches
|
return self.fit_loop.max_batches
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def num_sanity_val_batches(self) -> List[Union[int, float]]:
|
def num_sanity_val_batches(self) -> List[Union[int, float]]:
|
||||||
|
"""The number of validation batches that will be used during the sanity-checking part of
|
||||||
|
``trainer.fit()``."""
|
||||||
max_batches = self.fit_loop.epoch_loop.val_loop.max_batches
|
max_batches = self.fit_loop.epoch_loop.val_loop.max_batches
|
||||||
return [min(self.num_sanity_val_steps, batches) for batches in max_batches]
|
return [min(self.num_sanity_val_steps, batches) for batches in max_batches]
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def num_val_batches(self) -> List[Union[int, float]]:
|
def num_val_batches(self) -> List[Union[int, float]]:
|
||||||
|
"""The number of validation batches that will be used during ``trainer.fit()`` or
|
||||||
|
``trainer.validate()``."""
|
||||||
if self.state.fn == TrainerFn.VALIDATING:
|
if self.state.fn == TrainerFn.VALIDATING:
|
||||||
return self.validate_loop.max_batches
|
return self.validate_loop.max_batches
|
||||||
# if no trainer.fn is set, assume fit's validation
|
# if no trainer.fn is set, assume fit's validation
|
||||||
|
@ -1423,10 +1448,12 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def num_test_batches(self) -> List[Union[int, float]]:
|
def num_test_batches(self) -> List[Union[int, float]]:
|
||||||
|
"""The number of test batches that will be used during ``trainer.test()``."""
|
||||||
return self.test_loop.max_batches
|
return self.test_loop.max_batches
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def num_predict_batches(self) -> List[Union[int, float]]:
|
def num_predict_batches(self) -> List[Union[int, float]]:
|
||||||
|
"""The number of prediction batches that will be used during ``trainer.predict()``."""
|
||||||
return self.predict_loop.max_batches
|
return self.predict_loop.max_batches
|
||||||
|
|
||||||
@property
|
@property
|
||||||
|
@ -1454,6 +1481,7 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def logger(self) -> Optional[Logger]:
|
def logger(self) -> Optional[Logger]:
|
||||||
|
"""The first :class:`~lightning.pytorch.loggers.logger.Logger` being used."""
|
||||||
return self.loggers[0] if len(self.loggers) > 0 else None
|
return self.loggers[0] if len(self.loggers) > 0 else None
|
||||||
|
|
||||||
@logger.setter
|
@logger.setter
|
||||||
|
@ -1465,6 +1493,13 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def loggers(self) -> List[Logger]:
|
def loggers(self) -> List[Logger]:
|
||||||
|
"""The list of class:`~lightning.pytorch.loggers.logger.Logger` used.
|
||||||
|
|
||||||
|
..code-block:: python
|
||||||
|
|
||||||
|
for logger in trainer.loggers:
|
||||||
|
logger.log_metrics({"foo": 1.0})
|
||||||
|
"""
|
||||||
return self._loggers
|
return self._loggers
|
||||||
|
|
||||||
@loggers.setter
|
@loggers.setter
|
||||||
|
@ -1473,14 +1508,36 @@ class Trainer:
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def callback_metrics(self) -> _OUT_DICT:
|
def callback_metrics(self) -> _OUT_DICT:
|
||||||
|
"""The metrics available to callbacks.
|
||||||
|
|
||||||
|
This includes metrics logged via :meth:`~lightning.pytorch.core.module.LightningModule.log`.
|
||||||
|
|
||||||
|
..code-block:: python
|
||||||
|
|
||||||
|
def training_step(self, batch, batch_idx):
|
||||||
|
self.log("a_val", 2.0)
|
||||||
|
|
||||||
|
callback_metrics = trainer.callback_metrics
|
||||||
|
assert callback_metrics["a_val"] == 2.0
|
||||||
|
"""
|
||||||
return self._logger_connector.callback_metrics
|
return self._logger_connector.callback_metrics
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def logged_metrics(self) -> _OUT_DICT:
|
def logged_metrics(self) -> _OUT_DICT:
|
||||||
|
"""The metrics sent to the loggers.
|
||||||
|
|
||||||
|
This includes metrics logged via :meth:`~lightning.pytorch.core.module.LightningModule.log` with the
|
||||||
|
:paramref:`~lightning.pytorch.core.module.LightningModule.log.logger` argument set.
|
||||||
|
"""
|
||||||
return self._logger_connector.logged_metrics
|
return self._logger_connector.logged_metrics
|
||||||
|
|
||||||
@property
|
@property
|
||||||
def progress_bar_metrics(self) -> _PBAR_DICT:
|
def progress_bar_metrics(self) -> _PBAR_DICT:
|
||||||
|
"""The metrics sent to the progress bar.
|
||||||
|
|
||||||
|
This includes metrics logged via :meth:`~lightning.pytorch.core.module.LightningModule.log` with the
|
||||||
|
:paramref:`~lightning.pytorch.core.module.LightningModule.log.prog_bar` argument set.
|
||||||
|
"""
|
||||||
return self._logger_connector.progress_bar_metrics
|
return self._logger_connector.progress_bar_metrics
|
||||||
|
|
||||||
@property
|
@property
|
||||||
|
@ -1496,18 +1553,18 @@ class Trainer:
|
||||||
@property
|
@property
|
||||||
def estimated_stepping_batches(self) -> Union[int, float]:
|
def estimated_stepping_batches(self) -> Union[int, float]:
|
||||||
r"""
|
r"""
|
||||||
Estimated stepping batches for the complete training inferred from DataLoaders, gradient
|
The estimated number of batches that will ``optimizer.step()`` during training.
|
||||||
accumulation factor and distributed setup.
|
|
||||||
|
|
||||||
Examples::
|
This accounts for gradient accumulation and the current trainer configuration. This might sets up your training
|
||||||
|
dataloader if hadn't been set up already.
|
||||||
|
|
||||||
|
..code-block:: python
|
||||||
|
|
||||||
def configure_optimizers(self):
|
def configure_optimizers(self):
|
||||||
optimizer = ...
|
optimizer = ...
|
||||||
scheduler = torch.optim.lr_scheduler.OneCycleLR(
|
stepping_batches = self.trainer.estimated_stepping_batches
|
||||||
optimizer, max_lr=1e-3, total_steps=self.trainer.estimated_stepping_batches
|
scheduler = torch.optim.lr_scheduler.OneCycleLR(optimizer, max_lr=1e-3, total_steps=stepping_batches)
|
||||||
)
|
|
||||||
return [optimizer], [scheduler]
|
return [optimizer], [scheduler]
|
||||||
|
|
||||||
"""
|
"""
|
||||||
# infinite training
|
# infinite training
|
||||||
if self.max_epochs == -1:
|
if self.max_epochs == -1:
|
||||||
|
|
Loading…
Reference in New Issue