浏览代码
Enable generalization training (#2232)
Enable generalization training (#2232)
* Add Sampler and SamplerManager * Enable resampling of reset parameters during training * Documentation for Sampler and example YAML configuration file/develop-generalizationTraining-TrainerController
Ervin T
5 年前
当前提交
a46f3faa
共有 14 个文件被更改,包括 1857 次插入 和 43 次删除
-
4docs/Training-ML-Agents.md
-
8ml-agents-envs/mlagents/envs/exception.py
-
2ml-agents/mlagents/trainers/exception.py
-
67ml-agents/mlagents/trainers/learn.py
-
3ml-agents/mlagents/trainers/tests/test_environments/test_simple.py
-
13ml-agents/mlagents/trainers/tests/test_learn.py
-
19ml-agents/mlagents/trainers/tests/test_trainer_controller.py
-
81ml-agents/mlagents/trainers/trainer_controller.py
-
16config/generalize_test.yaml
-
124docs/Training-Generalization-Learning.md
-
850docs/images/3dball_big.png
-
482docs/images/3dball_small.png
-
134ml-agents-envs/mlagents/envs/sampler_class.py
-
97ml-agents-envs/mlagents/envs/tests/test_sampler_class.py
|
|||
resampling-interval: 5000 |
|||
|
|||
mass: |
|||
sampler-type: "uniform" |
|||
min_value: 0.5 |
|||
max_value: 10 |
|||
|
|||
gravity: |
|||
sampler-type: "uniform" |
|||
min_value: 7 |
|||
max_value: 12 |
|||
|
|||
scale: |
|||
sampler-type: "uniform" |
|||
min_value: 0.75 |
|||
max_value: 3 |
|
|||
# Training Generalized Reinforcement Learning Agents |
|||
|
|||
Reinforcement learning has a rather unique setup as opposed to supervised and |
|||
unsupervised learning. Agents here are trained and tested on the same exact |
|||
environment, which is analogous to a model being trained and tested on an |
|||
identical dataset in supervised learning! This setting results in overfitting; |
|||
the inability of the agent to generalize to slight tweaks or variations in the |
|||
environment. This is problematic in instances when environments are randomly |
|||
instantiated with varying properties. To make agents robust, one approach is to |
|||
train an agent over multiple variations of the environment. The agent is |
|||
trained in this approach with the intent that it learns to adapt its performance |
|||
to future unseen variations of the environment. |
|||
|
|||
Ball scale of 0.5 | Ball scale of 4 |
|||
:-------------------------:|:-------------------------: |
|||
![](images/3dball_small.png) | ![](images/3dball_big.png) |
|||
|
|||
_Variations of the 3D Ball environment._ |
|||
|
|||
To vary environments, we first decide what parameters to vary in an |
|||
environment. These parameters are known as `Reset Parameters`. In the 3D ball |
|||
environment example displayed in the figure above, the reset parameters are `gravity`, `ball_mass` and `ball_scale`. |
|||
|
|||
|
|||
## How-to |
|||
|
|||
For generalization training, we need to provide a way to modify the environment |
|||
by supplying a set of reset parameters, and vary them over time. This provision |
|||
can be done either deterministically or randomly. |
|||
|
|||
This is done by assigning each reset parameter a sampler, which samples a reset |
|||
parameter value (such as a uniform sampler). If a sampler isn't provided for a |
|||
reset parameter, the parameter maintains the default value throughout the |
|||
training, remaining unchanged. The samplers for all the reset parameters are |
|||
handled by a **Sampler Manager**, which also handles the generation of new |
|||
values for the reset parameters when needed. |
|||
|
|||
To setup the Sampler Manager, we setup a YAML file that specifies how we wish to |
|||
generate new samples. In this file, we specify the samplers and the |
|||
`resampling-duration` (number of simulation steps after which reset parameters are |
|||
resampled). Below is an example of a sampler file for the 3D ball environment. |
|||
|
|||
```yaml |
|||
episode-length: 5000 |
|||
|
|||
mass: |
|||
sampler-type: "uniform" |
|||
min_value: 0.5 |
|||
max_value: 10 |
|||
|
|||
gravity: |
|||
sampler-type: "multirange_uniform" |
|||
intervals: [[7, 10], [15, 20]] |
|||
|
|||
scale: |
|||
sampler-type: "uniform" |
|||
min_value: 0.75 |
|||
max_value: 3 |
|||
|
|||
``` |
|||
|
|||
* `resampling-duration` (int) - Specifies the number of steps for agent to |
|||
train under a particular environment configuration before resetting the |
|||
environment with a new sample of reset parameters. |
|||
|
|||
* `parameter_name` - Name of the reset parameter. This should match the name |
|||
specified in the academy of the intended environment for which the agent is |
|||
being trained. If a parameter specified in the file doesn't exist in the |
|||
environment, then this specification will be ignored. |
|||
|
|||
* `sampler-type` - Specify the sampler type to use for the reset parameter. |
|||
This is a string that should exist in the `Sampler Factory` (explained |
|||
below). |
|||
|
|||
* `sub-arguments` - Specify the characteristic parameters for the sampler. |
|||
In the example sampler file above, this would correspond to the `intervals` |
|||
key under the `multirange_uniform` sampler for the gravity reset parameter. |
|||
The key name should match the name of the corresponding argument in the sampler definition. (Look at defining a new sampler method) |
|||
|
|||
The sampler manager allocates a sampler for a reset parameter by using the *Sampler Factory*, which maintains a dictionary mapping of string keys to sampler objects. The available samplers to be used for reset parameter resampling is as available in the Sampler Factory. |
|||
|
|||
The implementation of the samplers can be found at `ml-agents-envs/mlagents/envs/sampler_class.py`. |
|||
|
|||
### Defining a new sampler method |
|||
|
|||
Custom sampling techniques must inherit from the *Sampler* base class (included in the `sampler_class` file) and preserve the interface. Once the class for the required method is specified, it must be registered in the Sampler Factory. |
|||
|
|||
This can be done by subscribing to the *register_sampler* method of the SamplerFactory. The command is as follows: |
|||
|
|||
`SamplerFactory.register_sampler(*custom_sampler_string_key*, *custom_sampler_object*)` |
|||
|
|||
Once the Sampler Factory reflects the new register, the custom sampler can be used for resampling reset parameter. For demonstration, lets say our sampler was implemented as below, and we register the `CustomSampler` class with the string `custom-sampler` in the Sampler Factory. |
|||
|
|||
```python |
|||
class CustomSampler(Sampler): |
|||
|
|||
def __init__(self, argA, argB, argC): |
|||
self.possible_vals = [argA, argB, argC] |
|||
|
|||
def sample_all(self): |
|||
return np.random.choice(self.possible_vals) |
|||
``` |
|||
|
|||
Now we need to specify this sampler in the sampler file. Lets say we wish to use this sampler for the reset parameter *mass*; the sampler file would specify the same for mass as the following (any order of the subarguments is valid). |
|||
|
|||
```yaml |
|||
mass: |
|||
sampler-type: "custom-sampler" |
|||
argB: 1 |
|||
argA: 2 |
|||
argC: 3 |
|||
``` |
|||
|
|||
With the sampler file setup, we can proceed to train our agent as explained in the next section. |
|||
|
|||
### Training with Generalization Learning |
|||
|
|||
We first begin with setting up the sampler file. After the sampler file is defined and configured, we proceed by launching `mlagents-learn` and specify our configured sampler file with the `--sampler` flag. To demonstrate, if we wanted to train a 3D ball agent with generalization using the `config/generalization-test.yaml` sampling setup, we can run |
|||
|
|||
```sh |
|||
mlagents-learn config/trainer_config.yaml --sampler=config/generalize_test.yaml --run-id=3D-Ball-generalization --train |
|||
``` |
|||
|
|||
We can observe progress and metrics via Tensorboard. |
|
|||
import numpy as np |
|||
from typing import * |
|||
from functools import * |
|||
from collections import OrderedDict |
|||
from abc import ABC, abstractmethod |
|||
|
|||
from .exception import SamplerException |
|||
|
|||
|
|||
class Sampler(ABC): |
|||
@abstractmethod |
|||
def sample_parameter(self) -> float: |
|||
pass |
|||
|
|||
|
|||
class UniformSampler(Sampler): |
|||
""" |
|||
Uniformly draws a single sample in the range [min_value, max_value). |
|||
""" |
|||
|
|||
def __init__( |
|||
self, min_value: Union[int, float], max_value: Union[int, float], **kwargs |
|||
) -> None: |
|||
self.min_value = min_value |
|||
self.max_value = max_value |
|||
|
|||
def sample_parameter(self) -> float: |
|||
return np.random.uniform(self.min_value, self.max_value) |
|||
|
|||
|
|||
class MultiRangeUniformSampler(Sampler): |
|||
""" |
|||
Draws a single sample uniformly from the intervals provided. The sampler |
|||
first picks an interval based on a weighted selection, with the weights |
|||
assigned to an interval based on its range. After picking the range, |
|||
it proceeds to pick a value uniformly in that range. |
|||
""" |
|||
|
|||
def __init__(self, intervals: List[List[Union[int, float]]], **kwargs) -> None: |
|||
self.intervals = intervals |
|||
# Measure the length of the intervals |
|||
interval_lengths = [abs(x[1] - x[0]) for x in self.intervals] |
|||
cum_interval_length = sum(interval_lengths) |
|||
# Assign weights to an interval proportionate to the interval size |
|||
self.interval_weights = [x / cum_interval_length for x in interval_lengths] |
|||
|
|||
def sample_parameter(self) -> float: |
|||
cur_min, cur_max = self.intervals[ |
|||
np.random.choice(len(self.intervals), p=self.interval_weights) |
|||
] |
|||
return np.random.uniform(cur_min, cur_max) |
|||
|
|||
|
|||
class GaussianSampler(Sampler): |
|||
""" |
|||
Draw a single sample value from a normal (gaussian) distribution. |
|||
This sampler is characterized by the mean and the standard deviation. |
|||
""" |
|||
|
|||
def __init__( |
|||
self, mean: Union[float, int], st_dev: Union[float, int], **kwargs |
|||
) -> None: |
|||
self.mean = mean |
|||
self.st_dev = st_dev |
|||
|
|||
def sample_parameter(self) -> float: |
|||
return np.random.normal(self.mean, self.st_dev) |
|||
|
|||
|
|||
class SamplerFactory: |
|||
""" |
|||
Maintain a directory of all samplers available. |
|||
Add new samplers using the register_sampler method. |
|||
""" |
|||
|
|||
NAME_TO_CLASS = { |
|||
"uniform": UniformSampler, |
|||
"gaussian": GaussianSampler, |
|||
"multirange_uniform": MultiRangeUniformSampler, |
|||
} |
|||
|
|||
@staticmethod |
|||
def register_sampler(name: str, sampler_cls: Type[Sampler]) -> None: |
|||
SamplerFactory.NAME_TO_CLASS[name] = sampler_cls |
|||
|
|||
@staticmethod |
|||
def init_sampler_class(name: str, params: Dict[str, Any]): |
|||
if name not in SamplerFactory.NAME_TO_CLASS: |
|||
raise SamplerException( |
|||
name + " sampler is not registered in the SamplerFactory." |
|||
" Use the register_sample method to register the string" |
|||
" associated to your sampler in the SamplerFactory." |
|||
) |
|||
sampler_cls = SamplerFactory.NAME_TO_CLASS[name] |
|||
try: |
|||
return sampler_cls(**params) |
|||
except TypeError: |
|||
raise SamplerException( |
|||
"The sampler class associated to the " + name + " key in the factory " |
|||
"was not provided the required arguments. Please ensure that the sampler " |
|||
"config file consists of the appropriate keys for this sampler class." |
|||
) |
|||
|
|||
|
|||
class SamplerManager: |
|||
def __init__(self, reset_param_dict: Dict[str, Any]) -> None: |
|||
self.reset_param_dict = reset_param_dict if reset_param_dict else {} |
|||
assert isinstance(self.reset_param_dict, dict) |
|||
self.samplers: Dict[str, Sampler] = {} |
|||
for param_name, cur_param_dict in self.reset_param_dict.items(): |
|||
if "sampler-type" not in cur_param_dict: |
|||
raise SamplerException( |
|||
"'sampler_type' argument hasn't been supplied for the {0} parameter".format( |
|||
param_name |
|||
) |
|||
) |
|||
sampler_name = cur_param_dict.pop("sampler-type") |
|||
param_sampler = SamplerFactory.init_sampler_class( |
|||
sampler_name, cur_param_dict |
|||
) |
|||
|
|||
self.samplers[param_name] = param_sampler |
|||
|
|||
def is_empty(self) -> bool: |
|||
""" |
|||
Check for if sampler_manager is empty. |
|||
""" |
|||
return not bool(self.samplers) |
|||
|
|||
def sample_all(self) -> Dict[str, float]: |
|||
res = {} |
|||
for param_name, param_sampler in list(self.samplers.items()): |
|||
res[param_name] = param_sampler.sample_parameter() |
|||
return res |
|
|||
from math import isclose |
|||
import pytest |
|||
|
|||
from mlagents.envs.sampler_class import SamplerManager |
|||
from mlagents.envs.sampler_class import ( |
|||
UniformSampler, |
|||
MultiRangeUniformSampler, |
|||
GaussianSampler, |
|||
) |
|||
from mlagents.envs.exception import UnityException |
|||
|
|||
|
|||
def sampler_config_1(): |
|||
return { |
|||
"mass": {"sampler-type": "uniform", "min_value": 5, "max_value": 10}, |
|||
"gravity": { |
|||
"sampler-type": "multirange_uniform", |
|||
"intervals": [[8, 11], [15, 20]], |
|||
}, |
|||
} |
|||
|
|||
|
|||
def check_value_in_intervals(val, intervals): |
|||
check_in_bounds = [a <= val <= b for a, b in intervals] |
|||
return any(check_in_bounds) |
|||
|
|||
|
|||
def test_sampler_config_1(): |
|||
config = sampler_config_1() |
|||
sampler = SamplerManager(config) |
|||
|
|||
assert sampler.is_empty() is False |
|||
assert isinstance(sampler.samplers["mass"], UniformSampler) |
|||
assert isinstance(sampler.samplers["gravity"], MultiRangeUniformSampler) |
|||
|
|||
cur_sample = sampler.sample_all() |
|||
|
|||
# Check uniform sampler for mass |
|||
assert sampler.samplers["mass"].min_value == config["mass"]["min_value"] |
|||
assert sampler.samplers["mass"].max_value == config["mass"]["max_value"] |
|||
assert config["mass"]["min_value"] <= cur_sample["mass"] |
|||
assert config["mass"]["max_value"] >= cur_sample["mass"] |
|||
|
|||
# Check multirange_uniform sampler for gravity |
|||
assert sampler.samplers["gravity"].intervals == config["gravity"]["intervals"] |
|||
assert check_value_in_intervals( |
|||
cur_sample["gravity"], sampler.samplers["gravity"].intervals |
|||
) |
|||
|
|||
|
|||
def sampler_config_2(): |
|||
return {"angle": {"sampler-type": "gaussian", "mean": 0, "st_dev": 1}} |
|||
|
|||
|
|||
def test_sampler_config_2(): |
|||
config = sampler_config_2() |
|||
sampler = SamplerManager(config) |
|||
assert sampler.is_empty() is False |
|||
assert isinstance(sampler.samplers["angle"], GaussianSampler) |
|||
|
|||
# Check angle gaussian sampler |
|||
assert sampler.samplers["angle"].mean == config["angle"]["mean"] |
|||
assert sampler.samplers["angle"].st_dev == config["angle"]["st_dev"] |
|||
|
|||
|
|||
def test_empty_samplers(): |
|||
empty_sampler = SamplerManager({}) |
|||
assert empty_sampler.is_empty() |
|||
empty_cur_sample = empty_sampler.sample_all() |
|||
assert empty_cur_sample == {} |
|||
|
|||
none_sampler = SamplerManager(None) |
|||
assert none_sampler.is_empty() |
|||
none_cur_sample = none_sampler.sample_all() |
|||
assert none_cur_sample == {} |
|||
|
|||
|
|||
def incorrect_uniform_sampler(): |
|||
# Do not specify required arguments to uniform sampler |
|||
return {"mass": {"sampler-type": "uniform", "min-value": 10}} |
|||
|
|||
|
|||
def incorrect_sampler_config(): |
|||
# Do not specify 'sampler-type' key |
|||
return {"mass": {"min-value": 2, "max-value": 30}} |
|||
|
|||
|
|||
def test_incorrect_uniform_sampler(): |
|||
config = incorrect_uniform_sampler() |
|||
with pytest.raises(UnityException): |
|||
SamplerManager(config) |
|||
|
|||
|
|||
def test_incorrect_sampler(): |
|||
config = incorrect_sampler_config() |
|||
with pytest.raises(UnityException): |
|||
SamplerManager(config) |
撰写
预览
正在加载...
取消
保存
Reference in new issue