Deterministic model

Deterministic models run continuous-domain deterministic policies.



skrl provides a Python mixin (DeterministicMixin) to assist in the creation of these types of models, allowing users to have full control over the function approximator definitions and architectures. Note that the use of this mixin must comply with the following rules:

  • The definition of multiple inheritance must always include the Model base class at the end.

  • The Model base class constructor must be invoked before the mixins constructor.

Warning

For models in JAX/Flax it is imperative to define all parameters (except observation_space, action_space and device) with default values to avoid errors (TypeError: __init__() missing N required positional argument) during initialization.

In addition, it is necessary to initialize the model’s state_dict (via the init_state_dict method) after its instantiation to avoid errors (AttributeError: object has no attribute "state_dict". If "state_dict" is defined in '.setup()', remember these fields are only accessible from inside 'init' or 'apply') during its use.

class DeterministicModel(DeterministicMixin, Model):
    def __init__(self, observation_space, action_space, device=None, clip_actions=False):
        Model.__init__(self, observation_space, action_space, device)
        DeterministicMixin.__init__(self, clip_actions)

Concept

Deterministic modelDeterministic model

Usage

  • Multi-Layer Perceptron (MLP)

  • Convolutional Neural Network (CNN)

  • Recurrent Neural Network (RNN)

  • Gated Recurrent Unit RNN (GRU)

  • Long Short-Term Memory RNN (LSTM)

../../_images/model_deterministic_mlp-light.svg../../_images/model_deterministic_mlp-dark.svg
import torch
import torch.nn as nn

from skrl.models.torch import Model, DeterministicMixin


# define the model
class MLP(DeterministicMixin, Model):
    def __init__(self, observation_space, action_space, device, clip_actions=False):
        Model.__init__(self, observation_space, action_space, device)
        DeterministicMixin.__init__(self, clip_actions)

        self.net = nn.Sequential(nn.Linear(self.num_observations + self.num_actions, 64),
                                 nn.ReLU(),
                                 nn.Linear(64, 32),
                                 nn.ReLU(),
                                 nn.Linear(32, 1))

    def compute(self, inputs, role):
        return self.net(torch.cat([inputs["states"], inputs["taken_actions"]], dim=1)), {}


# instantiate the model (assumes there is a wrapped environment: env)
critic = MLP(observation_space=env.observation_space,
             action_space=env.action_space,
             device=env.device,
             clip_actions=False)

API (PyTorch)

class skrl.models.torch.deterministic.DeterministicMixin(clip_actions: bool = False, role: str = '')

Bases: object

__init__(clip_actions: bool = False, role: str = '') None

Deterministic mixin model (deterministic model)

Parameters:
  • clip_actions (bool, optional) – Flag to indicate whether the actions should be clipped to the action space (default: False)

  • role (str, optional) – Role play by the model (default: "")

Example:

# define the model
>>> import torch
>>> import torch.nn as nn
>>> from skrl.models.torch import Model, DeterministicMixin
>>>
>>> class Value(DeterministicMixin, Model):
...     def __init__(self, observation_space, action_space, device="cuda:0", clip_actions=False):
...         Model.__init__(self, observation_space, action_space, device)
...         DeterministicMixin.__init__(self, clip_actions)
...
...         self.net = nn.Sequential(nn.Linear(self.num_observations, 32),
...                                  nn.ELU(),
...                                  nn.Linear(32, 32),
...                                  nn.ELU(),
...                                  nn.Linear(32, 1))
...
...     def compute(self, inputs, role):
...         return self.net(inputs["states"]), {}
...
>>> # given an observation_space: gym.spaces.Box with shape (60,)
>>> # and an action_space: gym.spaces.Box with shape (8,)
>>> model = Value(observation_space, action_space)
>>>
>>> print(model)
Value(
  (net): Sequential(
    (0): Linear(in_features=60, out_features=32, bias=True)
    (1): ELU(alpha=1.0)
    (2): Linear(in_features=32, out_features=32, bias=True)
    (3): ELU(alpha=1.0)
    (4): Linear(in_features=32, out_features=1, bias=True)
  )
)
act(inputs: Mapping[str, torch.Tensor | Any], role: str = '') Tuple[torch.Tensor, torch.Tensor | None, Mapping[str, torch.Tensor | Any]]

Act deterministically in response to the state of the environment

Parameters:
  • inputs (dict where the values are typically torch.Tensor) –

    Model inputs. The most common keys are:

    • "states": state of the environment used to make the decision

    • "taken_actions": actions taken by the policy for the given states

  • role (str, optional) – Role play by the model (default: "")

Returns:

Model output. The first component is the action to be taken by the agent. The second component is None. The third component is a dictionary containing extra output values

Return type:

tuple of torch.Tensor, torch.Tensor or None, and dict

Example:

>>> # given a batch of sample states with shape (4096, 60)
>>> actions, _, outputs = model.act({"states": states})
>>> print(actions.shape, outputs)
torch.Size([4096, 1]) {}

API (JAX)

class skrl.models.jax.deterministic.DeterministicMixin(clip_actions: bool = False, role: str = '')

Bases: object

__init__(clip_actions: bool = False, role: str = '') None

Deterministic mixin model (deterministic model)

Parameters:
  • clip_actions (bool, optional) – Flag to indicate whether the actions should be clipped to the action space (default: False)

  • role (str, optional) – Role play by the model (default: "")

Example:

# define the model
>>> import flax.linen as nn
>>> from skrl.models.jax import Model, DeterministicMixin
>>>
>>> class Value(DeterministicMixin, Model):
...     def __init__(self, observation_space, action_space, device=None, clip_actions=False, **kwargs):
...         Model.__init__(self, observation_space, action_space, device, **kwargs)
...         DeterministicMixin.__init__(self, clip_actions)
...
...     @nn.compact  # marks the given module method allowing inlined submodules
...     def __call__(self, inputs, role):
...         x = nn.elu(nn.Dense(32)([inputs["states"]))
...         x = nn.elu(nn.Dense(32)(x))
...         x = nn.Dense(1)(x)
...         return x, {}
...
>>> # given an observation_space: gym.spaces.Box with shape (60,)
>>> # and an action_space: gym.spaces.Box with shape (8,)
>>> model = Value(observation_space, action_space)
>>>
>>> print(model)
Value(
    # attributes
    observation_space = Box(-1.0, 1.0, (60,), float32)
    action_space = Box(-1.0, 1.0, (8,), float32)
    device = StreamExecutorGpuDevice(id=0, process_index=0, slice_index=0)
)
act(inputs: Mapping[str, ndarray | jax.Array | Any], role: str = '', params: jax.Array | None = None) Tuple[jax.Array, jax.Array | None, Mapping[str, jax.Array | Any]]

Act deterministically in response to the state of the environment

Parameters:
  • inputs (dict where the values are typically np.ndarray or jax.Array) –

    Model inputs. The most common keys are:

    • "states": state of the environment used to make the decision

    • "taken_actions": actions taken by the policy for the given states

  • role (str, optional) – Role play by the model (default: "")

  • params (jnp.array) – Parameters used to compute the output (default: None). If None, internal parameters will be used

Returns:

Model output. The first component is the action to be taken by the agent. The second component is None. The third component is a dictionary containing extra output values

Return type:

tuple of jax.Array, jax.Array or None, and dict

Example:

>>> # given a batch of sample states with shape (4096, 60)
>>> actions, _, outputs = model.act({"states": states})
>>> print(actions.shape, outputs)
(4096, 1) {}