Skip to content

Latest commit

 

History

History
552 lines (366 loc) · 26.7 KB

rllib-rlmodule.rst

File metadata and controls

552 lines (366 loc) · 26.7 KB

RL Modules (Alpha)

Note

This is an experimental module that serves as a general replacement for ModelV2, and is subject to change. It will eventually match the functionality of the previous stack. If you only use high-level RLlib APIs such as :py~ray.rllib.algorithms.algorithm.Algorithm you should not experience significant changes, except for a few new parameters to the configuration object. If you've used custom models or policies before, you'll need to migrate them to the new modules. Check the Migration guide for more information.

The table below shows the list of migrated algorithms and their current supported features, which will be updated as we progress.

Algorithm Independent MARL Fully-connected Image inputs (CNN) RNN support (LSTM) Complex observations (ComplexNet)
PPO pytorch tensorflow pytorch tensorflow pytorch tensorflow pytorch
Impala pytorch tensorflow pytorch tensorflow pytorch tensorflow pytorch
APPO pytorch tensorflow pytorch tensorflow pytorch tensorflow

RL Module is a neural network container that implements three public methods: :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_train, :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration, and :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference. Each method corresponds to a distinct reinforcement learning phase.

:py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration handles acting and data collection, balancing exploration and exploitation. On the other hand, the :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference serves the learned model during evaluation, often being less stochastic.

:py~ray.rllib.core.rl_module.rl_module.RLModule.forward_train manages the training phase, handling calculations exclusive to computing losses, such as learning Q values in a DQN model.

Enabling RL Modules in the Configuration

Enable RL Modules by setting the _enable_rl_module_api flag to True in the configuration object.

doc_code/rlmodule_guide.py

Constructing RL Modules

The RLModule API provides a unified way to define custom reinforcement learning models in RLlib. This API enables you to design and implement your own models to suit specific needs.

To maintain consistency and usability, RLlib offers a standardized approach for defining module objects for both single-agent and multi-agent reinforcement learning environments. This is achieved through the :py~ray.rllib.core.rl_module.rl_module.SingleAgentRLModuleSpec and :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModuleSpec classes. The built-in RLModules in RLlib follow this consistent design pattern, making it easier for you to understand and utilize these modules.

Single Agent

doc_code/rlmodule_guide.py

Multi Agent

doc_code/rlmodule_guide.py

You can pass RL Module specs to the algorithm configuration to be used by the algorithm.

Single Agent

doc_code/rlmodule_guide.py

Note

For passing RL Module specs, all fields do not have to be filled as they are filled based on the described environment or other algorithm configuration parameters (i.e. ,observation_space, action_space, model_config_dict are not required fields when passing a custom RL Module spec to the algorithm config.)

Multi Agent

doc_code/rlmodule_guide.py

Writing Custom Single Agent RL Modules

For single-agent algorithms (e.g., PPO, DQN) or independent multi-agent algorithms (e.g., PPO-MultiAgent), use :py~ray.rllib.core.rl_module.rl_module.RLModule. For more advanced multi-agent use cases with a shared communication between agents, extend the :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModule class.

RLlib treats single-agent modules as a special case of :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModule with only one module. Create the multi-agent representation of all RLModules by calling :py~ray.rllib.core.rl_module.rl_module.RLModule.as_multi_agent. For example:

doc_code/rlmodule_guide.py

RLlib implements the following abstract framework specific base classes:

  • TorchRLModule <ray.rllib.core.rl_module.torch_rl_module.TorchRLModule>: For PyTorch-based RL Modules.
  • TfRLModule <ray.rllib.core.rl_module.tf.tf_rl_module.TfRLModule>: For TensorFlow-based RL Modules.

The minimum requirement is for sub-classes of :py~ray.rllib.core.rl_module.rl_module.RLModule is to implement the following methods:

  • :py~ray.rllib.core.rl_module.rl_module.RLModule._forward_train: Forward pass for training.
  • :py~ray.rllib.core.rl_module.rl_module.RLModule._forward_inference: Forward pass for inference.
  • :py~ray.rllib.core.rl_module.rl_module.RLModule._forward_exploration: Forward pass for exploration.

For your custom :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration and :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference methods, you must return a dictionary that either contains the key "actions" and/or the key "action_dist_inputs".

If you return the "actions" key:

  • RLlib will use the actions provided thereunder as-is.
  • If you also returned the "action_dist_inputs" key: RLlib will also create a :py~ray.rllib.models.distributions.Distribution object from the distribution parameters under that key and - in the case of :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration - compute action probs and logp values from the given actions automatically.

If you do not return the "actions" key:

  • You must return the "action_dist_inputs" key instead from your :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration and :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference methods.
  • RLlib will create a :py~ray.rllib.models.distributions.Distribution object from the distribution parameters under that key and sample actions from the thus generated distribution.
  • In the case of :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration, RLlib will also compute action probs and logp values from the sampled actions automatically.

Note

In the case of :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference, the generated distributions (from returned key "action_dist_inputs") will always be made deterministic first via the :py~ray.rllib.models.distributions.Distribution.to_deterministic utility before a possible action sample step. Thus, for example, sampling from a Categorical distribution will be reduced to simply selecting the argmax actions from the distribution's logits/probs.

Commonly used distribution implementations can be found under ray.rllib.models.tf.tf_distributions for tensorflow and ray.rllib.models.torch.torch_distributions for torch. You can choose to return determinstic actions, by creating a determinstic distribution instance.

Returning "actions" key

"""
An RLModule whose forward_exploration/inference methods return the
"actions" key.
"""

class MyRLModule(TorchRLModule):
    ...

    def _forward_inference(self, batch):
        ...
        return {
            "actions": ...  # actions will be used as-is
        }

    def _forward_exploration(self, batch):
        ...
        return {
            "actions": ...  # actions will be used as-is (no sampling step!)
            "action_dist_inputs": ...  # optional: If provided, will be used to compute action probs and logp.
        }

Not returning "actions" key

"""
An RLModule whose forward_exploration/inference methods do NOT return the
"actions" key.
"""

class MyRLModule(TorchRLModule):
    ...

    def _forward_inference(self, batch):
        ...
        return {
            # RLlib will:
            # - Generate distribution from these parameters.
            # - Convert distribution to a deterministic equivalent.
            # - "sample" from the deterministic distribution.
            "action_dist_inputs": ...
        }

    def _forward_exploration(self, batch):
        ...
        return {
            # RLlib will:
            # - Generate distribution from these parameters.
            # - "sample" from the (stochastic) distribution.
            # - Compute action probs/logs automatically using the sampled
            #   actions and the generated distribution object.
            "action_dist_inputs": ...
        }

Also the :py~ray.rllib.core.rl_module.rl_module.RLModule class's constrcutor requires a dataclass config object called ~ray.rllib.core.rl_module.rl_module.RLModuleConfig which contains the following fields:

  • :py~ray.rllib.core.rl_module.rl_module.RLModuleConfig.observation_space: The observation space of the environment (either processed or raw).
  • :py~ray.rllib.core.rl_module.rl_module.RLModuleConfig.action_space: The action space of the environment.
  • :py~ray.rllib.core.rl_module.rl_module.RLModuleConfig.model_config_dict: The model config dictionary of the algorithm. Model hyper-parameters such as number of layers, type of activation, etc. are defined here.
  • :py~ray.rllib.core.rl_module.rl_module.RLModuleConfig.catalog_class: The :py~ray.rllib.core.models.catalog.Catalog object of the algorithm.

When writing RL Modules, you need to use these fields to construct your model.

Single Agent (torch)

doc_code/rlmodule_guide.py

Single Agent (tensorflow)

doc_code/rlmodule_guide.py

In :py~ray.rllib.core.rl_module.rl_module.RLModule you can enforce the checking for the existence of certain input or output keys in the data that is communicated into and out of RL Modules. This serves multiple purposes:

  • For the I/O requirement of each method to be self-documenting.
  • For failures to happen quickly. If users extend the modules and implement something that does not match the assumptions of the I/O specs, the check reports missing keys and their expected format. For example, RLModule should always have an obs key in the input batch and an action_dist key in the output.

Single Level Keys

doc_code/rlmodule_guide.py

Nested Keys

doc_code/rlmodule_guide.py

TensorShape Spec

doc_code/rlmodule_guide.py

Type Spec

doc_code/rlmodule_guide.py

:py~ray.rllib.core.rl_module.rl_module.RLModule has two methods for each forward method, totaling 6 methods that can be override to describe the specs of the input and output of each method:

  • :py~ray.rllib.core.rl_module.rl_module.RLModule.input_specs_inference
  • :py~ray.rllib.core.rl_module.rl_module.RLModule.output_specs_inference
  • :py~ray.rllib.core.rl_module.rl_module.RLModule.input_specs_exploration
  • :py~ray.rllib.core.rl_module.rl_module.RLModule.output_specs_exploration
  • :py~ray.rllib.core.rl_module.rl_module.RLModule.input_specs_train
  • :py~ray.rllib.core.rl_module.rl_module.RLModule.output_specs_train

To learn more, see the SpecType documentation.

Writing Custom Multi-Agent RL Modules (Advanced)

For multi-agent modules, RLlib implements :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModule, which is a dictionary of :py~ray.rllib.core.rl_module.rl_module.RLModule objects, one for each policy, and possibly some shared modules. The base-class implementation works for most of use cases that need to define independent neural networks for sub-groups of agents. For more complex, multi-agent use cases, where the agents share some part of their neural network, you should inherit from this class and override the default implementation.

The :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModule offers an API for constructing custom models tailored to specific needs. The key method for this customization is :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModule.build.

The following example creates a custom multi-agent RL module with underlying modules. The modules share an encoder, which gets applied to the global part of the observations space. The local part passes through a separate encoder, specific to each policy.

Multi agent with shared encoder (Torch)

doc_code/rlmodule_guide.py

To construct this custom multi-agent RL module, pass the class to the :py~ray.rllib.core.rl_module.marl_module.MultiAgentRLModuleSpec constructor. Also, pass the :py~ray.rllib.core.rl_module.rl_module.SingleAgentRLModuleSpec for each agent because RLlib requires the observation, action spaces, and model hyper-parameters for each agent.

doc_code/rlmodule_guide.py

Extending Existing RLlib RL Modules

RLlib provides a number of RL Modules for different frameworks (e.g., PyTorch, TensorFlow, etc.). Extend these modules by inheriting from them and overriding the methods you need to customize. For example, extend :py~ray.rllib.algorithms.ppo.torch.ppo_torch_rl_module.PPOTorchRLModule and augment it with your own customization. Then pass the new customized class into the algorithm configuration.

There are two possible ways to extend existing RL Modules:

Inheriting existing RL Modules

One way to extend existing RL Modules is to inherit from them and override the methods you need to customize. For example, extend :py~ray.rllib.algorithms.ppo.torch.ppo_torch_rl_module.PPOTorchRLModule and augment it with your own customization. Then pass the new customized class into the algorithm configuration to use the PPO algorithm to optimize your custom RL Module.

class MyPPORLModule(PPORLModule):

    def __init__(self, config: RLModuleConfig):
        super().__init__(config)
        ...

# Pass in the custom RL Module class to the spec
algo_config = algo_config.rl_module(
    rl_module_spec=SingleAgentRLModuleSpec(module_class=MyPPORLModule)
)

Extending RL Module Catalog

Another way to customize your module is by extending its :py~ray.rllib.core.models.catalog.Catalog. The :py~ray.rllib.core.models.catalog.Catalog is a component that defines the default architecture and behavior of a model based on factors such as observation_space, action_space, etc. To modify sub-components of an existing RL Module, extend the corresponding Catalog class.

For instance, to adapt the existing PPORLModule for a custom graph observation space not supported by RLlib out-of-the-box, extend the :py~ray.rllib.core.models.catalog.Catalog class used to create the PPORLModule and override the method responsible for returning the encoder component to ensure that your custom encoder replaces the default one initially provided by RLlib. For more information on the :py~ray.rllib.core.models.catalog.Catalog class, refer to the Catalog user guide.

class MyAwesomeCatalog(PPOCatalog):

    def get_actor_critic_encoder_config():
        # create your awesome graph encoder here and return it
        pass


# Pass in the custom catalog class to the spec
algo_config = algo_config.rl_module(
    rl_module_spec=SingleAgentRLModuleSpec(catalog_class=MyAwesomeCatalog)
)

Migrating from Custom Policies and Models to RL Modules

This document is for those who have implemented custom policies and models in RLlib and want to migrate to the new ~ray.rllib.core.rl_module.rl_module.RLModule API. If you have implemented custom policies that extended the ~ray.rllib.policy.eager_tf_policy_v2.EagerTFPolicyV2 or ~ray.rllib.policy.torch_policy_v2.TorchPolicyV2 classes, you likely did so that you could either modify the behavior of constructing models and distributions (via overriding ~ray.rllib.policy.torch_policy_v2.TorchPolicyV2.make_model, ~ray.rllib.policy.torch_policy_v2.TorchPolicyV2.make_model_and_action_dist), control the action sampling logic (via overriding ~ray.rllib.policy.eager_tf_policy_v2.EagerTFPolicyV2.action_distribution_fn or ~ray.rllib.policy.eager_tf_policy_v2.EagerTFPolicyV2.action_sampler_fn), or control the logic for infernce (via overriding ~ray.rllib.policy.policy.Policy.compute_actions_from_input_dict, ~ray.rllib.policy.policy.Policy.compute_actions, or ~ray.rllib.policy.policy.Policy.compute_log_likelihoods). These APIs were built with ray.rllib.models.modelv2.ModelV2 models in mind to enable you to customize the behavior of those functions. However ~ray.rllib.core.rl_module.rl_module.RLModule is a more general abstraction that will reduce the amount of functions that you need to override.

In the new ~ray.rllib.core.rl_module.rl_module.RLModule API the construction of the models and the action distribution class that should be used are best defined in the constructor. That RL Module is constructed automatically if users follow the instructions outlined in the sections Enabling RL Modules in the Configuration and Constructing RL Modules. ~ray.rllib.policy.policy.Policy.compute_actions and ~ray.rllib.policy.policy.Policy.compute_actions_from_input_dict can still be used for sampling actions for inference or exploration by using the explore=True|False parameter. If called with explore=True these functions will invoke ~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration and if explore=False then they will call ~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference.

What your customization could have looked like before:

ModelV2

from ray.rllib.models.torch.torch_modelv2 import TorchModelV2
from ray.rllib.policy.torch_policy_v2 import TorchPolicyV2


class MyCustomModel(TorchModelV2):
    """Code for your previous custom model"""
    ...


class CustomPolicy(TorchPolicyV2):

    @DeveloperAPI
    @OverrideToImplementCustomLogic
    def make_model(self) -> ModelV2:
        """Create model.

        Note: only one of make_model or make_model_and_action_dist
        can be overridden.

        Returns:
        ModelV2 model.
        """
        return MyCustomModel(...)

ModelV2 + Distribution

from ray.rllib.models.torch.torch_modelv2 import TorchModelV2
from ray.rllib.policy.torch_policy_v2 import TorchPolicyV2


class MyCustomModel(TorchModelV2):
    """Code for your previous custom model"""
    ...


class CustomPolicy(TorchPolicyV2):

    @DeveloperAPI
    @OverrideToImplementCustomLogic
    def make_model_and_action_dist(self):
        """Create model and action distribution function.

        Returns:
            ModelV2 model.
            ActionDistribution class.
        """
        my_model = MyCustomModel(...) # construct some ModelV2 instance here
        dist_class = ... # Action distribution cls

        return my_model, dist_class

Sampler functions

from ray.rllib.models.torch.torch_modelv2 import TorchModelV2
from ray.rllib.policy.torch_policy_v2 import TorchPolicyV2

class CustomPolicy(TorchPolicyV2):

    @DeveloperAPI
    @OverrideToImplementCustomLogic
    def action_sampler_fn(
        self,
        model: ModelV2,
        *,
        obs_batch: TensorType,
        state_batches: TensorType,
        **kwargs,
    ) -> Tuple[TensorType, TensorType, TensorType, List[TensorType]]:
        """Custom function for sampling new actions given policy.

        Args:
            model: Underlying model.
            obs_batch: Observation tensor batch.
            state_batches: Action sampling state batch.

        Returns:
            Sampled action
            Log-likelihood
            Action distribution inputs
            Updated state
        """
        return None, None, None, None


    @DeveloperAPI
    @OverrideToImplementCustomLogic
    def action_distribution_fn(
        self,
        model: ModelV2,
        *,
        obs_batch: TensorType,
        state_batches: TensorType,
        **kwargs,
    ) -> Tuple[TensorType, type, List[TensorType]]:
        """Action distribution function for this Policy.

        Args:
            model: Underlying model.
            obs_batch: Observation tensor batch.
            state_batches: Action sampling state batch.

        Returns:
            Distribution input.
            ActionDistribution class.
            State outs.
        """
        return None, None, None

All of the Policy.compute_*** functions expect that :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_exploration and :py~ray.rllib.core.rl_module.rl_module.RLModule.forward_inference return a dictionary that either contains the key "actions" and/or the key "action_dist_inputs".

See Writing Custom Single Agent RL Modules for more details on how to implement your own custom RL Modules.

The Equivalent RL Module

"""
No need to override any policy functions. Simply instead implement any custom logic in your custom RL Module
"""
from ray.rllib.models.torch.torch_distributions import YOUR_DIST_CLASS


class MyRLModule(TorchRLModule):

    def __init__(self, config: RLConfig):
        # construct any custom networks here using config
        # specify an action distribution class here
        ...

    def _forward_inference(self, batch):
        ...

    def _forward_exploration(self, batch):
        ...

Notable TODOs

  • [] Add support for RNNs.
  • [] Checkpointing.
  • [] End to end example for custom RL Modules extending PPORLModule (e.g. LLM)